2023-12-08 16:29:13 +00:00
|
|
|
#include <c10/util/Flags.h>
|
2021-07-26 18:51:41 +00:00
|
|
|
#include <c10/util/irange.h>
|
2020-03-07 17:59:11 +00:00
|
|
|
#include <torch/csrc/jit/api/function_impl.h>
|
2020-03-26 18:04:28 +00:00
|
|
|
#include <torch/csrc/jit/passes/inliner.h>
|
2020-03-26 20:27:07 +00:00
|
|
|
|
|
|
|
|
#include <torch/csrc/jit/frontend/error_report.h>
|
2020-04-07 16:39:56 +00:00
|
|
|
#include <torch/csrc/jit/passes/constant_pooling.h>
|
|
|
|
|
#include <torch/csrc/jit/passes/constant_propagation.h>
|
|
|
|
|
#include <torch/csrc/jit/passes/peephole.h>
|
2019-06-10 23:31:18 +00:00
|
|
|
|
Add fp16/fp32 autocasting to JIT/TorchScript (#63939)
Summary:
Adds mixed precision autocasting support between fp32/fp16 to torchscript/JIT. More in depth descriptoin can be found at [torch/csrc/jit/JIT-AUTOCAST.md](https://github.com/pytorch/pytorch/pull/63939/files#diff-1f1772aaa508841c5bb58b74ab98f49a1e577612cd9ea5c386c8714a75db830b)
This PR implemented an autocast optimization pass that inserts casting ops per AMP rule (torch/csrc/jit/passes/autocast.cpp), that mimics the behavior of eager autocast. The pass also takes into consideration the context of `torch.cuda.amp.autocast` and only inserts casting ops within the enabled context manager, giving feature parity as with eager amp autocast.
We currently provide JIT AMP autocast as a prototyping feature, so it is default off and could be turned on via `torch._C._jit_set_autocast_mode(True)`
The JIT support for autocast is subject to different constraints compared to the eager mode implementation (mostly related to the fact that TorchScript is statically typed), restriction on the user facing python code is described in doc torch/csrc/jit/JIT-AUTOCAST.md
This is a prototype, there are also implementation limitation that's necessary to keep this PR small and get something functioning quickly on upstream, so we can iterate on designs.
Few limitation/challenge that is not properly resolved in this PR:
1. Autocast inserts cast operation, which would have impact on scalar type of output tensor feeding downstream operations. We are not currently propagating the updated scalar types, this would give issues/wrong results on operations in promotion rules.
2. Backward for autodiff in JIT misses the casting of dgrad to input scalar type, as what autograd does in eager. This forces us to explicitly mark the casting operation for certain operations (e.g. binary ops), otherwise, we might be feeding dgrad with mismatch scalar type to input. This could potentially break gradient function consuming dgrad. (e.g. gemm backwards, which assumes grad_output to be of same scalar type as input')
3. `torch.autocast` api has an optional argument `dtype` which is not currently supported in the JIT autocast and we require a static value.
Credit goes mostly to:
tlemo
kevinstephano
Pull Request resolved: https://github.com/pytorch/pytorch/pull/63939
Reviewed By: navahgar
Differential Revision: D31093381
Pulled By: eellison
fbshipit-source-id: da6e26c668c38b01e296f304507048d6c1794314
2021-10-27 19:09:53 +00:00
|
|
|
#ifndef C10_MOBILE
|
|
|
|
|
#include <ATen/autocast_mode.h>
|
|
|
|
|
#include <torch/csrc/jit/passes/autocast.h>
|
|
|
|
|
#endif
|
|
|
|
|
|
2023-12-08 16:29:13 +00:00
|
|
|
C10_DEFINE_bool(
|
|
|
|
|
torch_jit_do_not_store_optimized_graph,
|
|
|
|
|
false,
|
|
|
|
|
"Do not store the optimized graph.");
|
|
|
|
|
|
2023-01-13 16:32:30 +00:00
|
|
|
namespace torch::jit {
|
2019-08-23 22:56:46 +00:00
|
|
|
namespace {
|
2021-10-27 18:52:48 +00:00
|
|
|
c10::FunctionSchema defaultSchemaFor(const GraphFunction& function) {
|
2020-03-07 17:59:11 +00:00
|
|
|
std::vector<c10::Argument> args;
|
|
|
|
|
std::vector<c10::Argument> returns;
|
2019-08-23 22:56:46 +00:00
|
|
|
Graph& g = *function.graph();
|
|
|
|
|
size_t num_inputs = function.num_inputs();
|
2021-07-26 18:51:41 +00:00
|
|
|
for (const auto i : c10::irange(num_inputs)) {
|
2019-08-23 22:56:46 +00:00
|
|
|
const Value* v = g.inputs().at(i);
|
|
|
|
|
std::string name = v->hasDebugName() ? v->debugNameBase()
|
2024-06-10 23:40:45 +00:00
|
|
|
: ("argument_" + std::to_string(i));
|
2019-08-23 22:56:46 +00:00
|
|
|
args.emplace_back(std::move(name), unshapedType(g.inputs()[i]->type()));
|
|
|
|
|
}
|
2021-07-26 18:51:41 +00:00
|
|
|
for (const auto i : c10::irange(g.outputs().size())) {
|
2019-08-23 22:56:46 +00:00
|
|
|
returns.emplace_back("", unshapedType(g.outputs()[i]->type()));
|
|
|
|
|
}
|
|
|
|
|
return {function.name(), "", std::move(args), std::move(returns)};
|
|
|
|
|
}
|
2021-10-27 18:52:48 +00:00
|
|
|
|
|
|
|
|
template <typename T, typename F>
|
|
|
|
|
T* tryToGraphFunctionImpl(F& function) noexcept {
|
|
|
|
|
if (!function.isGraphFunction()) {
|
|
|
|
|
return nullptr;
|
|
|
|
|
}
|
|
|
|
|
|
|
|
|
|
return static_cast<T*>(&function);
|
|
|
|
|
}
|
|
|
|
|
|
|
|
|
|
template <typename T, typename F>
|
|
|
|
|
T& toGraphFunctionImpl(F& function) {
|
|
|
|
|
if (auto* g = tryToGraphFunctionImpl<T>(function)) {
|
|
|
|
|
return *g;
|
|
|
|
|
}
|
|
|
|
|
|
|
|
|
|
TORCH_INTERNAL_ASSERT(
|
|
|
|
|
false,
|
|
|
|
|
"Failed to downcast a Function to a GraphFunction. "
|
|
|
|
|
"This probably indicates that the JIT calling context needs a "
|
|
|
|
|
"special case on tryToGraphFunction() instead.");
|
|
|
|
|
}
|
|
|
|
|
|
2019-08-23 22:56:46 +00:00
|
|
|
} // namespace
|
|
|
|
|
|
2023-06-02 22:04:39 +00:00
|
|
|
static void placeholderCreator(GraphFunction&) {
|
2019-06-10 23:31:18 +00:00
|
|
|
throw RecursiveMethodCallError();
|
|
|
|
|
}
|
|
|
|
|
|
2020-03-07 17:59:11 +00:00
|
|
|
void GraphFunction::run(Stack& stack) {
|
Bytecode export flow (#25187)
Summary:
Pull Request resolved: https://github.com/pytorch/pytorch/pull/25187
The bytecode export flow: dump the bytecode format for the light weighted interpreter.
* The bytecode is generated without input spec optimization. It would be more generic (input independent) with no obvious performance degradation (to be tested).
* Main API: torch::jit::script::Module::save(filename, extra_files, bool *bytecode_format* = false).
* Both bytecode and module object are exported in pickle format.
* The module object (in data.pkl) is the same as the original JIT model.
* The serializer is dependent on pickle only (no protobuf or Json).
* The major functionality is forked in ScriptModuleSerializer2::serialize().
* The test loader is test_bc_export.cpp.
* Simple APIs are added in Code and its implementation to get necessary information (instructions, operators and constants).
* Since there's no dependency on graph/node, GetAttr is promoted from an operator to first-class instruction (https://github.com/pytorch/pytorch/pull/25151) .
* Some definitions (instructions, writeArchive, etc) that are shared by full JIT and bytecode are pulled out of the local namespace (https://github.com/pytorch/pytorch/pull/25148).
The output layout looks like:
* folders of methods.
* In each method folder (for example, forward/):
* bytecode.pkl: instructions and operators
* constants{.pkl,/}: constant list in constants.pkl. If there are tensors in constants, the binary tensor files in constants/ folder.
* data{.pkl,/}: the module object, with binary tensor files in data/ folder. The same as in torchscript.
Test Plan: Imported from OSS
Differential Revision: D17076411
fbshipit-source-id: 46eb298e7320d1e585b0101effc0fcfd09219046
2019-09-25 23:34:05 +00:00
|
|
|
get_executor().run(stack);
|
|
|
|
|
}
|
|
|
|
|
|
2020-11-05 01:06:05 +00:00
|
|
|
c10::intrusive_ptr<c10::ivalue::Future> GraphFunction::runAsync(
|
|
|
|
|
Stack& stack,
|
|
|
|
|
TaskLauncher taskLauncher) {
|
|
|
|
|
return get_executor().runAsync(stack, std::move(taskLauncher));
|
2020-04-01 00:19:33 +00:00
|
|
|
}
|
|
|
|
|
|
2020-03-07 17:59:11 +00:00
|
|
|
void GraphFunction::ensure_defined() {
|
2020-02-19 07:42:46 +00:00
|
|
|
if (function_creator_) {
|
|
|
|
|
auto creator = function_creator_;
|
|
|
|
|
function_creator_ = placeholderCreator;
|
|
|
|
|
creator(*this);
|
|
|
|
|
function_creator_ = nullptr;
|
2019-06-10 23:31:18 +00:00
|
|
|
}
|
|
|
|
|
check_single_output();
|
|
|
|
|
}
|
2019-07-05 00:07:52 +00:00
|
|
|
|
2020-03-07 17:59:11 +00:00
|
|
|
const c10::FunctionSchema& GraphFunction::getSchema() const {
|
2019-08-23 22:56:46 +00:00
|
|
|
if (schema_ == nullptr) {
|
2020-03-07 17:59:11 +00:00
|
|
|
schema_ = std::make_unique<c10::FunctionSchema>(defaultSchemaFor(*this));
|
2019-08-23 22:56:46 +00:00
|
|
|
}
|
|
|
|
|
return *schema_;
|
|
|
|
|
}
|
2019-09-24 07:20:15 +00:00
|
|
|
|
2023-12-08 16:29:13 +00:00
|
|
|
std::shared_ptr<Graph> GraphFunction::optimized_graph() const {
|
|
|
|
|
std::lock_guard<std::recursive_mutex> lock(compile_mutex);
|
2023-12-09 20:43:37 +00:00
|
|
|
decltype(optimized_graphs_)::value_type graph;
|
|
|
|
|
auto& graph_ref = !FLAGS_torch_jit_do_not_store_optimized_graph
|
2023-12-08 16:29:13 +00:00
|
|
|
? optimized_graphs_[currentSpecialization()]
|
2023-12-09 20:43:37 +00:00
|
|
|
: graph;
|
|
|
|
|
if (graph_ref) {
|
|
|
|
|
return graph_ref;
|
2023-12-08 16:29:13 +00:00
|
|
|
}
|
2023-12-09 20:43:37 +00:00
|
|
|
graph_ref = graph_->copy();
|
2023-12-08 16:29:13 +00:00
|
|
|
if (getGraphExecutorOptimize()) {
|
2023-12-09 20:43:37 +00:00
|
|
|
preoptimizeGraph(graph_ref, force_no_amp_);
|
2023-12-08 16:29:13 +00:00
|
|
|
}
|
2023-12-09 20:43:37 +00:00
|
|
|
return graph_ref;
|
2023-12-08 16:29:13 +00:00
|
|
|
}
|
|
|
|
|
|
Add fp16/fp32 autocasting to JIT/TorchScript (#63939)
Summary:
Adds mixed precision autocasting support between fp32/fp16 to torchscript/JIT. More in depth descriptoin can be found at [torch/csrc/jit/JIT-AUTOCAST.md](https://github.com/pytorch/pytorch/pull/63939/files#diff-1f1772aaa508841c5bb58b74ab98f49a1e577612cd9ea5c386c8714a75db830b)
This PR implemented an autocast optimization pass that inserts casting ops per AMP rule (torch/csrc/jit/passes/autocast.cpp), that mimics the behavior of eager autocast. The pass also takes into consideration the context of `torch.cuda.amp.autocast` and only inserts casting ops within the enabled context manager, giving feature parity as with eager amp autocast.
We currently provide JIT AMP autocast as a prototyping feature, so it is default off and could be turned on via `torch._C._jit_set_autocast_mode(True)`
The JIT support for autocast is subject to different constraints compared to the eager mode implementation (mostly related to the fact that TorchScript is statically typed), restriction on the user facing python code is described in doc torch/csrc/jit/JIT-AUTOCAST.md
This is a prototype, there are also implementation limitation that's necessary to keep this PR small and get something functioning quickly on upstream, so we can iterate on designs.
Few limitation/challenge that is not properly resolved in this PR:
1. Autocast inserts cast operation, which would have impact on scalar type of output tensor feeding downstream operations. We are not currently propagating the updated scalar types, this would give issues/wrong results on operations in promotion rules.
2. Backward for autodiff in JIT misses the casting of dgrad to input scalar type, as what autograd does in eager. This forces us to explicitly mark the casting operation for certain operations (e.g. binary ops), otherwise, we might be feeding dgrad with mismatch scalar type to input. This could potentially break gradient function consuming dgrad. (e.g. gemm backwards, which assumes grad_output to be of same scalar type as input')
3. `torch.autocast` api has an optional argument `dtype` which is not currently supported in the JIT autocast and we require a static value.
Credit goes mostly to:
tlemo
kevinstephano
Pull Request resolved: https://github.com/pytorch/pytorch/pull/63939
Reviewed By: navahgar
Differential Revision: D31093381
Pulled By: eellison
fbshipit-source-id: da6e26c668c38b01e296f304507048d6c1794314
2021-10-27 19:09:53 +00:00
|
|
|
GraphFunction::SpecializationKey GraphFunction::currentSpecialization() const {
|
2022-05-16 17:14:09 +00:00
|
|
|
if (force_no_amp_) {
|
|
|
|
|
return SpecializationKey::AutocastOff;
|
|
|
|
|
}
|
Add fp16/fp32 autocasting to JIT/TorchScript (#63939)
Summary:
Adds mixed precision autocasting support between fp32/fp16 to torchscript/JIT. More in depth descriptoin can be found at [torch/csrc/jit/JIT-AUTOCAST.md](https://github.com/pytorch/pytorch/pull/63939/files#diff-1f1772aaa508841c5bb58b74ab98f49a1e577612cd9ea5c386c8714a75db830b)
This PR implemented an autocast optimization pass that inserts casting ops per AMP rule (torch/csrc/jit/passes/autocast.cpp), that mimics the behavior of eager autocast. The pass also takes into consideration the context of `torch.cuda.amp.autocast` and only inserts casting ops within the enabled context manager, giving feature parity as with eager amp autocast.
We currently provide JIT AMP autocast as a prototyping feature, so it is default off and could be turned on via `torch._C._jit_set_autocast_mode(True)`
The JIT support for autocast is subject to different constraints compared to the eager mode implementation (mostly related to the fact that TorchScript is statically typed), restriction on the user facing python code is described in doc torch/csrc/jit/JIT-AUTOCAST.md
This is a prototype, there are also implementation limitation that's necessary to keep this PR small and get something functioning quickly on upstream, so we can iterate on designs.
Few limitation/challenge that is not properly resolved in this PR:
1. Autocast inserts cast operation, which would have impact on scalar type of output tensor feeding downstream operations. We are not currently propagating the updated scalar types, this would give issues/wrong results on operations in promotion rules.
2. Backward for autodiff in JIT misses the casting of dgrad to input scalar type, as what autograd does in eager. This forces us to explicitly mark the casting operation for certain operations (e.g. binary ops), otherwise, we might be feeding dgrad with mismatch scalar type to input. This could potentially break gradient function consuming dgrad. (e.g. gemm backwards, which assumes grad_output to be of same scalar type as input')
3. `torch.autocast` api has an optional argument `dtype` which is not currently supported in the JIT autocast and we require a static value.
Credit goes mostly to:
tlemo
kevinstephano
Pull Request resolved: https://github.com/pytorch/pytorch/pull/63939
Reviewed By: navahgar
Differential Revision: D31093381
Pulled By: eellison
fbshipit-source-id: da6e26c668c38b01e296f304507048d6c1794314
2021-10-27 19:09:53 +00:00
|
|
|
#ifdef C10_MOBILE
|
|
|
|
|
// disabling autodiff pass for mobile build since autocast APIs don't exist
|
|
|
|
|
return SpecializationKey::AutocastOff;
|
|
|
|
|
#else
|
2024-04-23 09:34:45 +00:00
|
|
|
bool cpu_enabled = at::autocast::is_autocast_enabled(at::kCPU);
|
|
|
|
|
bool gpu_enabled = at::autocast::is_autocast_enabled(at::kCUDA);
|
Add fp16/fp32 autocasting to JIT/TorchScript (#63939)
Summary:
Adds mixed precision autocasting support between fp32/fp16 to torchscript/JIT. More in depth descriptoin can be found at [torch/csrc/jit/JIT-AUTOCAST.md](https://github.com/pytorch/pytorch/pull/63939/files#diff-1f1772aaa508841c5bb58b74ab98f49a1e577612cd9ea5c386c8714a75db830b)
This PR implemented an autocast optimization pass that inserts casting ops per AMP rule (torch/csrc/jit/passes/autocast.cpp), that mimics the behavior of eager autocast. The pass also takes into consideration the context of `torch.cuda.amp.autocast` and only inserts casting ops within the enabled context manager, giving feature parity as with eager amp autocast.
We currently provide JIT AMP autocast as a prototyping feature, so it is default off and could be turned on via `torch._C._jit_set_autocast_mode(True)`
The JIT support for autocast is subject to different constraints compared to the eager mode implementation (mostly related to the fact that TorchScript is statically typed), restriction on the user facing python code is described in doc torch/csrc/jit/JIT-AUTOCAST.md
This is a prototype, there are also implementation limitation that's necessary to keep this PR small and get something functioning quickly on upstream, so we can iterate on designs.
Few limitation/challenge that is not properly resolved in this PR:
1. Autocast inserts cast operation, which would have impact on scalar type of output tensor feeding downstream operations. We are not currently propagating the updated scalar types, this would give issues/wrong results on operations in promotion rules.
2. Backward for autodiff in JIT misses the casting of dgrad to input scalar type, as what autograd does in eager. This forces us to explicitly mark the casting operation for certain operations (e.g. binary ops), otherwise, we might be feeding dgrad with mismatch scalar type to input. This could potentially break gradient function consuming dgrad. (e.g. gemm backwards, which assumes grad_output to be of same scalar type as input')
3. `torch.autocast` api has an optional argument `dtype` which is not currently supported in the JIT autocast and we require a static value.
Credit goes mostly to:
tlemo
kevinstephano
Pull Request resolved: https://github.com/pytorch/pytorch/pull/63939
Reviewed By: navahgar
Differential Revision: D31093381
Pulled By: eellison
fbshipit-source-id: da6e26c668c38b01e296f304507048d6c1794314
2021-10-27 19:09:53 +00:00
|
|
|
if (cpu_enabled && gpu_enabled) {
|
|
|
|
|
return SpecializationKey::CpuGpuAutocastOn;
|
|
|
|
|
} else if (!cpu_enabled && !gpu_enabled) {
|
|
|
|
|
return SpecializationKey::AutocastOff;
|
|
|
|
|
} else {
|
|
|
|
|
return gpu_enabled ? SpecializationKey::GpuAutocastOn
|
|
|
|
|
: SpecializationKey::CpuAutocastOn;
|
|
|
|
|
}
|
|
|
|
|
#endif
|
|
|
|
|
}
|
|
|
|
|
|
2022-05-16 17:14:09 +00:00
|
|
|
void preoptimizeGraph(std::shared_ptr<Graph>& graph, bool disable_autocast) {
|
2019-09-24 07:20:15 +00:00
|
|
|
Inline(*graph);
|
Add fp16/fp32 autocasting to JIT/TorchScript (#63939)
Summary:
Adds mixed precision autocasting support between fp32/fp16 to torchscript/JIT. More in depth descriptoin can be found at [torch/csrc/jit/JIT-AUTOCAST.md](https://github.com/pytorch/pytorch/pull/63939/files#diff-1f1772aaa508841c5bb58b74ab98f49a1e577612cd9ea5c386c8714a75db830b)
This PR implemented an autocast optimization pass that inserts casting ops per AMP rule (torch/csrc/jit/passes/autocast.cpp), that mimics the behavior of eager autocast. The pass also takes into consideration the context of `torch.cuda.amp.autocast` and only inserts casting ops within the enabled context manager, giving feature parity as with eager amp autocast.
We currently provide JIT AMP autocast as a prototyping feature, so it is default off and could be turned on via `torch._C._jit_set_autocast_mode(True)`
The JIT support for autocast is subject to different constraints compared to the eager mode implementation (mostly related to the fact that TorchScript is statically typed), restriction on the user facing python code is described in doc torch/csrc/jit/JIT-AUTOCAST.md
This is a prototype, there are also implementation limitation that's necessary to keep this PR small and get something functioning quickly on upstream, so we can iterate on designs.
Few limitation/challenge that is not properly resolved in this PR:
1. Autocast inserts cast operation, which would have impact on scalar type of output tensor feeding downstream operations. We are not currently propagating the updated scalar types, this would give issues/wrong results on operations in promotion rules.
2. Backward for autodiff in JIT misses the casting of dgrad to input scalar type, as what autograd does in eager. This forces us to explicitly mark the casting operation for certain operations (e.g. binary ops), otherwise, we might be feeding dgrad with mismatch scalar type to input. This could potentially break gradient function consuming dgrad. (e.g. gemm backwards, which assumes grad_output to be of same scalar type as input')
3. `torch.autocast` api has an optional argument `dtype` which is not currently supported in the JIT autocast and we require a static value.
Credit goes mostly to:
tlemo
kevinstephano
Pull Request resolved: https://github.com/pytorch/pytorch/pull/63939
Reviewed By: navahgar
Differential Revision: D31093381
Pulled By: eellison
fbshipit-source-id: da6e26c668c38b01e296f304507048d6c1794314
2021-10-27 19:09:53 +00:00
|
|
|
|
2020-04-07 16:39:56 +00:00
|
|
|
// Peephole Optimize cleans up many "is None" checks and creates constant prop
|
|
|
|
|
// opportunities
|
2020-04-16 00:32:25 +00:00
|
|
|
PeepholeOptimize(graph, true);
|
Add fp16/fp32 autocasting to JIT/TorchScript (#63939)
Summary:
Adds mixed precision autocasting support between fp32/fp16 to torchscript/JIT. More in depth descriptoin can be found at [torch/csrc/jit/JIT-AUTOCAST.md](https://github.com/pytorch/pytorch/pull/63939/files#diff-1f1772aaa508841c5bb58b74ab98f49a1e577612cd9ea5c386c8714a75db830b)
This PR implemented an autocast optimization pass that inserts casting ops per AMP rule (torch/csrc/jit/passes/autocast.cpp), that mimics the behavior of eager autocast. The pass also takes into consideration the context of `torch.cuda.amp.autocast` and only inserts casting ops within the enabled context manager, giving feature parity as with eager amp autocast.
We currently provide JIT AMP autocast as a prototyping feature, so it is default off and could be turned on via `torch._C._jit_set_autocast_mode(True)`
The JIT support for autocast is subject to different constraints compared to the eager mode implementation (mostly related to the fact that TorchScript is statically typed), restriction on the user facing python code is described in doc torch/csrc/jit/JIT-AUTOCAST.md
This is a prototype, there are also implementation limitation that's necessary to keep this PR small and get something functioning quickly on upstream, so we can iterate on designs.
Few limitation/challenge that is not properly resolved in this PR:
1. Autocast inserts cast operation, which would have impact on scalar type of output tensor feeding downstream operations. We are not currently propagating the updated scalar types, this would give issues/wrong results on operations in promotion rules.
2. Backward for autodiff in JIT misses the casting of dgrad to input scalar type, as what autograd does in eager. This forces us to explicitly mark the casting operation for certain operations (e.g. binary ops), otherwise, we might be feeding dgrad with mismatch scalar type to input. This could potentially break gradient function consuming dgrad. (e.g. gemm backwards, which assumes grad_output to be of same scalar type as input')
3. `torch.autocast` api has an optional argument `dtype` which is not currently supported in the JIT autocast and we require a static value.
Credit goes mostly to:
tlemo
kevinstephano
Pull Request resolved: https://github.com/pytorch/pytorch/pull/63939
Reviewed By: navahgar
Differential Revision: D31093381
Pulled By: eellison
fbshipit-source-id: da6e26c668c38b01e296f304507048d6c1794314
2021-10-27 19:09:53 +00:00
|
|
|
|
|
|
|
|
// AliasDb construction can be slow, so run it just on immutable types
|
|
|
|
|
// to clean up constant Ifs & other easy wins
|
2020-04-07 16:39:56 +00:00
|
|
|
ConstantPropagationImmutableTypes(graph);
|
Add fp16/fp32 autocasting to JIT/TorchScript (#63939)
Summary:
Adds mixed precision autocasting support between fp32/fp16 to torchscript/JIT. More in depth descriptoin can be found at [torch/csrc/jit/JIT-AUTOCAST.md](https://github.com/pytorch/pytorch/pull/63939/files#diff-1f1772aaa508841c5bb58b74ab98f49a1e577612cd9ea5c386c8714a75db830b)
This PR implemented an autocast optimization pass that inserts casting ops per AMP rule (torch/csrc/jit/passes/autocast.cpp), that mimics the behavior of eager autocast. The pass also takes into consideration the context of `torch.cuda.amp.autocast` and only inserts casting ops within the enabled context manager, giving feature parity as with eager amp autocast.
We currently provide JIT AMP autocast as a prototyping feature, so it is default off and could be turned on via `torch._C._jit_set_autocast_mode(True)`
The JIT support for autocast is subject to different constraints compared to the eager mode implementation (mostly related to the fact that TorchScript is statically typed), restriction on the user facing python code is described in doc torch/csrc/jit/JIT-AUTOCAST.md
This is a prototype, there are also implementation limitation that's necessary to keep this PR small and get something functioning quickly on upstream, so we can iterate on designs.
Few limitation/challenge that is not properly resolved in this PR:
1. Autocast inserts cast operation, which would have impact on scalar type of output tensor feeding downstream operations. We are not currently propagating the updated scalar types, this would give issues/wrong results on operations in promotion rules.
2. Backward for autodiff in JIT misses the casting of dgrad to input scalar type, as what autograd does in eager. This forces us to explicitly mark the casting operation for certain operations (e.g. binary ops), otherwise, we might be feeding dgrad with mismatch scalar type to input. This could potentially break gradient function consuming dgrad. (e.g. gemm backwards, which assumes grad_output to be of same scalar type as input')
3. `torch.autocast` api has an optional argument `dtype` which is not currently supported in the JIT autocast and we require a static value.
Credit goes mostly to:
tlemo
kevinstephano
Pull Request resolved: https://github.com/pytorch/pytorch/pull/63939
Reviewed By: navahgar
Differential Revision: D31093381
Pulled By: eellison
fbshipit-source-id: da6e26c668c38b01e296f304507048d6c1794314
2021-10-27 19:09:53 +00:00
|
|
|
|
|
|
|
|
#ifndef C10_MOBILE
|
|
|
|
|
// Inject casts for automatic mixed precision
|
|
|
|
|
//
|
|
|
|
|
// TODO: Ideally, this pass could run earlier, before inlining
|
|
|
|
|
// or any other optimizations. That setup is preferable because:
|
|
|
|
|
// 1. The AMP pass would be self-contained and function independently
|
|
|
|
|
// of the any optimizations
|
|
|
|
|
// 2. AMP transformations would benefit from followup passes's cleanup
|
|
|
|
|
//
|
2022-05-16 17:14:09 +00:00
|
|
|
if (!disable_autocast) {
|
|
|
|
|
Autocast(graph);
|
|
|
|
|
}
|
Add fp16/fp32 autocasting to JIT/TorchScript (#63939)
Summary:
Adds mixed precision autocasting support between fp32/fp16 to torchscript/JIT. More in depth descriptoin can be found at [torch/csrc/jit/JIT-AUTOCAST.md](https://github.com/pytorch/pytorch/pull/63939/files#diff-1f1772aaa508841c5bb58b74ab98f49a1e577612cd9ea5c386c8714a75db830b)
This PR implemented an autocast optimization pass that inserts casting ops per AMP rule (torch/csrc/jit/passes/autocast.cpp), that mimics the behavior of eager autocast. The pass also takes into consideration the context of `torch.cuda.amp.autocast` and only inserts casting ops within the enabled context manager, giving feature parity as with eager amp autocast.
We currently provide JIT AMP autocast as a prototyping feature, so it is default off and could be turned on via `torch._C._jit_set_autocast_mode(True)`
The JIT support for autocast is subject to different constraints compared to the eager mode implementation (mostly related to the fact that TorchScript is statically typed), restriction on the user facing python code is described in doc torch/csrc/jit/JIT-AUTOCAST.md
This is a prototype, there are also implementation limitation that's necessary to keep this PR small and get something functioning quickly on upstream, so we can iterate on designs.
Few limitation/challenge that is not properly resolved in this PR:
1. Autocast inserts cast operation, which would have impact on scalar type of output tensor feeding downstream operations. We are not currently propagating the updated scalar types, this would give issues/wrong results on operations in promotion rules.
2. Backward for autodiff in JIT misses the casting of dgrad to input scalar type, as what autograd does in eager. This forces us to explicitly mark the casting operation for certain operations (e.g. binary ops), otherwise, we might be feeding dgrad with mismatch scalar type to input. This could potentially break gradient function consuming dgrad. (e.g. gemm backwards, which assumes grad_output to be of same scalar type as input')
3. `torch.autocast` api has an optional argument `dtype` which is not currently supported in the JIT autocast and we require a static value.
Credit goes mostly to:
tlemo
kevinstephano
Pull Request resolved: https://github.com/pytorch/pytorch/pull/63939
Reviewed By: navahgar
Differential Revision: D31093381
Pulled By: eellison
fbshipit-source-id: da6e26c668c38b01e296f304507048d6c1794314
2021-10-27 19:09:53 +00:00
|
|
|
#endif
|
|
|
|
|
|
2020-04-07 16:39:56 +00:00
|
|
|
ConstantPooling(graph);
|
2019-09-24 07:20:15 +00:00
|
|
|
}
|
|
|
|
|
|
2021-10-27 18:52:48 +00:00
|
|
|
GraphFunction* tryToGraphFunction(Function& function) noexcept {
|
|
|
|
|
return tryToGraphFunctionImpl<GraphFunction>(function);
|
|
|
|
|
}
|
|
|
|
|
|
|
|
|
|
GraphFunction& toGraphFunction(Function& function) {
|
|
|
|
|
return toGraphFunctionImpl<GraphFunction>(function);
|
|
|
|
|
}
|
|
|
|
|
|
|
|
|
|
const GraphFunction& toGraphFunction(const Function& function) {
|
|
|
|
|
return toGraphFunctionImpl<const GraphFunction>(function);
|
|
|
|
|
}
|
|
|
|
|
|
2023-01-13 16:32:30 +00:00
|
|
|
} // namespace torch::jit
|