forked from pytorch/pytorch
-
Notifications
You must be signed in to change notification settings - Fork 0
/
tracer.h
415 lines (363 loc) · 12.6 KB
/
tracer.h
1
2
3
4
5
6
7
8
9
10
11
12
13
14
15
16
17
18
19
20
21
22
23
24
25
26
27
28
29
30
31
32
33
34
35
36
37
38
39
40
41
42
43
44
45
46
47
48
49
50
51
52
53
54
55
56
57
58
59
60
61
62
63
64
65
66
67
68
69
70
71
72
73
74
75
76
77
78
79
80
81
82
83
84
85
86
87
88
89
90
91
92
93
94
95
96
97
98
99
100
101
102
103
104
105
106
107
108
109
110
111
112
113
114
115
116
117
118
119
120
121
122
123
124
125
126
127
128
129
130
131
132
133
134
135
136
137
138
139
140
141
142
143
144
145
146
147
148
149
150
151
152
153
154
155
156
157
158
159
160
161
162
163
164
165
166
167
168
169
170
171
172
173
174
175
176
177
178
179
180
181
182
183
184
185
186
187
188
189
190
191
192
193
194
195
196
197
198
199
200
201
202
203
204
205
206
207
208
209
210
211
212
213
214
215
216
217
218
219
220
221
222
223
224
225
226
227
228
229
230
231
232
233
234
235
236
237
238
239
240
241
242
243
244
245
246
247
248
249
250
251
252
253
254
255
256
257
258
259
260
261
262
263
264
265
266
267
268
269
270
271
272
273
274
275
276
277
278
279
280
281
282
283
284
285
286
287
288
289
290
291
292
293
294
295
296
297
298
299
300
301
302
303
304
305
306
307
308
309
310
311
312
313
314
315
316
317
318
319
320
321
322
323
324
325
326
327
328
329
330
331
332
333
334
335
336
337
338
339
340
341
342
343
344
345
346
347
348
349
350
351
352
353
354
355
356
357
358
359
360
361
362
363
364
365
366
367
368
369
370
371
372
373
374
375
376
377
378
379
380
381
382
383
384
385
386
387
388
389
390
391
392
393
394
395
396
397
398
399
400
401
402
403
404
405
406
407
408
409
410
411
412
413
414
415
#pragma once
#include <ATen/core/Dimname.h>
#include <ATen/core/class_type.h>
#include <ATen/core/jit_type.h>
#include <ATen/core/stack.h>
#include <ATen/core/symbol.h>
#include <c10/util/Exception.h>
#include <torch/csrc/Export.h>
#include <torch/csrc/jit/frontend/source_range.h>
#include <torch/csrc/utils/variadic.h>
#include <cstdint>
#include <iostream>
#include <memory>
#include <mutex>
#include <unordered_map>
#include <vector>
namespace torch {
namespace jit {
struct Node;
struct Value;
struct Graph;
struct Module;
namespace tracer {
using ::c10::ivalue::Shared;
using ::c10::IValue;
using ::c10::ivalue::Future;
using ::c10::ArrayRef;
using ::c10::TupleType;
using ::c10::TupleTypePtr;
using ::c10::ivalue::ConstantString;
using torch::autograd::Variable;
using variable_list = std::vector<Variable>;
TORCH_API std::atomic<bool>& getTracerStateWarnMode();
struct TORCH_API TracingState
: public std::enable_shared_from_this<TracingState> {
TracingState();
~TracingState();
// NOLINTNEXTLINE(cppcoreguidelines-non-private-member-variables-in-classes)
std::shared_ptr<Graph> graph;
// NOLINTNEXTLINE(cppcoreguidelines-non-private-member-variables-in-classes)
bool warn = getTracerStateWarnMode();
// NOLINTNEXTLINE(cppcoreguidelines-non-private-member-variables-in-classes)
bool strict = true;
// NOLINTNEXTLINE(cppcoreguidelines-non-private-member-variables-in-classes)
bool force_outplace = false;
// NOLINTNEXTLINE(cppcoreguidelines-non-private-member-variables-in-classes)
std::function<std::string(const Variable& var)> lookup_var_name_fn =
[](const Variable& var) { return ""; };
void enterFrame() {
env_stack.emplace_back();
}
void leaveFrame() {
env_stack.pop_back();
}
void setValue(const IValue& v, Value* value);
void delValue(const IValue& var);
Value* getValue(const IValue& var);
Value* getOutput(const IValue& var, size_t i);
bool hasValue(const IValue& var) const;
Node* createNode(c10::Symbol op_name, size_t num_outputs);
void insertNode(Node* node);
private:
using WeakIValue = at::WeakIValue;
struct WeakIValueHasher {
size_t operator()(const WeakIValue& t) const {
return t.hash();
}
};
struct WeakIValueEq {
bool operator()(const WeakIValue& t1, const WeakIValue& t2) const {
return t1.isSameIdentity(t2);
}
};
using Frame =
std::unordered_map<WeakIValue, Value*, WeakIValueHasher, WeakIValueEq>;
std::vector<Frame> env_stack;
};
// This is meant to be used as a thread local place, where we can store extra
// info that gets lost when we call into ATen from Python bindings. One example
// for when this happens is when we get an IntArrayRef argument with e.g. sizes
// for view. When tracing, those might be tensors, which let us encode extra
// data dependencies, but once they get to the ATen call where we actually have
// the tracing logic, they get converted into a raw IntArrayRef, and we loose
// all information. To prevent this, we temporarily stash it in here.
// NOLINTNEXTLINE(cppcoreguidelines-pro-type-member-init)
struct ArgumentStash {
struct IntArrayRefTrace : std::vector<Value*> {
IntArrayRefTrace(int size) : std::vector<Value*>(size, nullptr) {}
};
static bool empty() {
return stash.intlists.empty();
}
TORCH_API static void stashIntArrayRefElem(
const std::string& arg_name,
size_t size,
size_t idx,
const Variable& var);
static bool hasIntArrayRef(const std::string& arg_name) {
return stash.intlists.count(arg_name) > 0;
}
static IntArrayRefTrace popIntArrayRef(const std::string& arg_name) {
auto info = std::move(stash.intlists.at(arg_name));
stash.intlists.erase(arg_name);
return info;
}
// Value stashing: Use these methods to stash arguments which correspond
// to regular Value*'s in the graph. i.e. they don't require special
// handling like in the case of IntArrayRefs
TORCH_API static void stashValue(
const std::string& arg_name,
size_t idx,
const Variable& var,
const c10::TypePtr& type = nullptr);
static bool hasValue(const std::string& arg_name) {
return stash.values.count(arg_name) > 0;
}
static Value* popValue(const std::string& arg_name) {
auto info = stash.values.at(arg_name);
stash.values.erase(arg_name);
return info;
}
private:
static thread_local ArgumentStash stash;
std::unordered_map<std::string, IntArrayRefTrace> intlists;
std::unordered_map<std::string, Value*> values;
};
// Retrieve or set the current tracing state. Returns a nullptr if tracing is
// disabled.
TORCH_API const std::shared_ptr<TracingState>& getTracingState();
TORCH_API void setTracingState(std::shared_ptr<TracingState> state);
inline bool isTracing() {
return static_cast<bool>(getTracingState());
}
using warn_fn_type = void (*)(const std::string& msg);
TORCH_API extern const char* WARN_PYTHON_DATAFLOW;
TORCH_API extern const char* WARN_CONSTRUCTOR;
TORCH_API extern const char* WARN_RESIZE;
TORCH_API extern const char* STRICT_TRACER_MSG;
TORCH_API void _do_warn(const char* _reason, const char* _kind);
inline void warn(const char* _reason, const char* _kind = nullptr) {
if (const auto& state = getTracingState()) {
if (!state->warn)
return;
_do_warn(_reason, _kind);
}
}
TORCH_API void setWarn(warn_fn_type fn);
struct TORCH_API NoWarn {
NoWarn() : state(getTracingState()) {
if (state) {
prev = state->warn;
state->warn = false;
}
}
~NoWarn() {
if (state) {
state->warn = prev;
}
}
std::shared_ptr<TracingState> state;
bool prev{false};
};
struct WithNestedTracingFrame {
WithNestedTracingFrame() {
getTracingState()->enterFrame();
}
~WithNestedTracingFrame() {
getTracingState()->leaveFrame();
}
};
TORCH_API void recordSourceLocation(Node* n);
TORCH_API void setRecordSourceLocation(void (*v)(Node*));
TORCH_API std::vector<StackEntry> pythonCallstack();
TORCH_API void setPythonCallstack(std::vector<StackEntry> (*v)());
// Having finished adding a new 'node' to the graph IR 'setValueTrace'
// associates this node with an output variable, so that further operations
// involving this variable know which node in the IR to reference.
TORCH_API void setValueTrace(const IValue& v, Value* value);
TORCH_API void delValueTrace(const IValue& var);
TORCH_API std::function<void()> pauseTracing();
TORCH_API Value* getValueTrace(const IValue& var);
TORCH_API std::pair<std::shared_ptr<TracingState>, Stack> trace(
Stack inputs,
const std::function<Stack(Stack)>& traced_fn,
std::function<std::string(const Variable&)> var_name_lookup_fn,
bool strict = true,
bool force_outplace = false,
Module* self = nullptr,
const std::vector<std::string>& argument_names = {});
TORCH_API void abandon();
// NB: those serve both as an intermediate steps in addInputs below,
// as well as the overloads that terminate template recursion
TORCH_API void addInputs(Node* n, const char* name, int64_t value);
TORCH_API void addInputs(Node* n, const char* name, c10::SymInt value);
TORCH_API void addInputs(
Node* n,
const char* name,
c10::optional<int64_t> value);
TORCH_API void addInputs(Node* n, const char* name, bool value);
TORCH_API void addInputs(
Node* n,
const char* name,
const c10::optional<bool>& value);
TORCH_API void addInputs(Node* n, const char* name, double value);
TORCH_API void addInputs(
Node* n,
const char* name,
const c10::optional<double>& value);
TORCH_API void addInputs(Node* n, const char* name, const at::Scalar& value);
TORCH_API void addInputs(
Node* n,
const char* name,
const c10::optional<at::Scalar>& value);
TORCH_API void addInputs(Node* n, const char* name, const at::Tensor& value);
TORCH_API void addInputs(
Node* n,
const char* name,
const c10::optional<at::Tensor>& value);
TORCH_API void addInputs(Node* n, const char* name, ArrayRef<int64_t> value);
TORCH_API void addInputs(Node* n, const char* name, c10::SymIntArrayRef value);
TORCH_API void addInputs(
Node* n,
const char* name,
c10::optional<c10::SymInt> value);
TORCH_API void addInputs(
Node* n,
const char* name,
const c10::optional<ArrayRef<int64_t>>& value);
TORCH_API void addInputs(
Node* n,
const char* name,
const at::OptionalIntArrayRef& opt_value);
TORCH_API void addInputs(
Node* n,
const char* name,
const at::OptionalSymIntArrayRef& opt_value);
TORCH_API void addInputs(
Node* n,
const char* name,
ArrayRef<at::Tensor> value,
bool allow_undefined = false);
TORCH_API void addInputs(
Node* n,
const char* name,
std::vector<at::Tensor> value,
bool allow_undefined = false);
TORCH_API void addInputs(
Node* n,
const char* name,
at::ITensorListRef value,
bool allow_undefined = false);
TORCH_API void addInputs(
Node* n,
const char* name,
const List<c10::optional<at::Tensor>>& value);
TORCH_API void addInputs(
Node* n,
const char* name,
ArrayRef<c10::intrusive_ptr<c10::ivalue::Object>> value,
const c10::ClassTypePtr& class_type);
TORCH_API void addInputs(Node* n, const char* name, ArrayRef<double> value);
TORCH_API void addInputs(
Node* n,
const char* name,
const c10::optional<ArrayRef<double>>& value);
TORCH_API void addInputs(
Node* n,
const char* name,
const c10::string_view value);
TORCH_API void addInputs(
Node* n,
const char* name,
const c10::optional<c10::string_view>& value);
TORCH_API void addInputs(Node* n, const char* name, at::Device value);
TORCH_API void addInputs(Node* n, const char* name, c10::Stream stream);
TORCH_API void addInputs(Node* n, const char* name, at::Layout value);
TORCH_API void addInputs(Node* n, const char* name, at::ScalarType value);
TORCH_API void addInputs(
Node* n,
const char* name,
const c10::optional<at::ScalarType>& value);
TORCH_API void addInputs(
Node* n,
const char* name,
const c10::optional<at::Device>& value);
TORCH_API void addInputs(
Node* n,
const char* name,
const c10::optional<at::Layout>& value);
TORCH_API void addInputs(Node* n, const char* name, at::MemoryFormat value);
TORCH_API void addInputs(
Node* n,
const char* name,
c10::optional<at::DimnameList> value);
TORCH_API void addInputs(
Node* n,
const char* name,
const c10::optional<at::MemoryFormat>& value);
TORCH_API void addInputs(
Node* n,
const char* name,
const c10::optional<at::Generator>& value);
inline void addInputs(
Node* n,
const char* name,
const std::vector<bool>& value) {
AT_ERROR("Tracing a list of bool type is currently not supported!");
}
template <typename T>
void addInputs(Node* n, const char* name, ArrayRef<T> value) {
AT_ERROR("Tracing a list of arbitrary type is currently not supported!");
}
template <typename K, typename V>
void addInputs(
Node* n,
const char* name,
const std::unordered_map<K, V>& value) {
AT_ERROR("Tracing a dict of arbitrary types is currently not supported!");
}
template <size_t N>
void addInputs(Node* n, const char* name, std::array<bool, N> value) {
throw std::runtime_error(
"Found an unsupported argument type in the JIT tracer. File a bug report.");
}
TORCH_API void addInputs(
Node* n,
const char* name,
const c10::intrusive_ptr<c10::ivalue::Object>& obj);
TORCH_API void ensureUniqueIfOutOfPlaced(
const char* name,
const at::Tensor& tensor);
TORCH_API void ensureUniqueIfOutOfPlaced(
const char* name,
const c10::optional<at::Tensor>& tensor);
template <
typename T,
typename = torch::enable_if_t<(
!std::is_convertible<torch::decay_t<T>, at::TensorList>::value &&
!std::is_convertible<torch::decay_t<T>, c10::List<at::Tensor>>::value &&
!std::is_convertible<torch::decay_t<T>, at::Tensor>::value &&
!std::is_convertible<
torch::decay_t<T>,
c10::intrusive_ptr<c10::ivalue::Object>>::value)>>
void addOutput(Node* node, T&&) {
AT_ERROR(
"Found an unsupported argument type ",
c10::demangle_type<T>(),
" in the JIT tracer. File a bug report.");
}
TORCH_API void addOutput(Node* node, const at::Tensor& tensor);
TORCH_API void setOutput(Value* value, const at::Tensor& output);
TORCH_API void addOutput(Node* node, const std::vector<at::Tensor>& list);
TORCH_API void addOutput(Node* node, const c10::List<at::Tensor>& list);
TORCH_API void addOutput(
Node* node,
const c10::intrusive_ptr<c10::ivalue::Object>& output);
TORCH_API autograd::Variable getSizeOf(
const autograd::Variable& var,
int64_t dim);
TORCH_API autograd::Variable getNumelOf(const autograd::Variable& var);
} // namespace tracer
} // namespace jit
} // namespace torch