Skip to content

Commit fc037c1

Browse files
authored
Merge pull request #1892 from pytorch/fb-sync-shirong
Sync fb internal change to OSS
2 parents d4e5ed0 + f77a002 commit fc037c1

26 files changed

+1193
-197
lines changed

py/torch_tensorrt/fx/converters/acc_ops_converters.py

+5-2
Original file line numberDiff line numberDiff line change
@@ -691,10 +691,13 @@ def acc_ops_layer_norm(network, target, args, kwargs, name):
691691
eps_field = trt.PluginField(
692692
"eps", np.array([kwargs["eps"]], dtype=np.float32), trt.PluginFieldType.FLOAT32
693693
)
694+
normalized_shape = kwargs["normalized_shape"]
694695
try:
695-
normalized_shape = np.array(kwargs["normalized_shape"], dtype=np.int32)
696+
normalized_shape = np.array(normalized_shape, dtype=np.int32)
696697
except TypeError:
697-
_LOGGER.error("Unable to convert normalized_shape to a field, fall back to []")
698+
_LOGGER.error(
699+
f"Unable to convert normalized_shape with value {normalized_shape} to a field, fall back to []"
700+
)
698701
normalized_shape = np.array([], dtype=np.int32)
699702

700703
normalized_shape_filed = trt.PluginField(

py/torch_tensorrt/fx/diagnostics.py

+8
Original file line numberDiff line numberDiff line change
@@ -87,12 +87,14 @@ class DiagnosticsWriter:
8787

8888
def __init__(self):
8989
self._root_dir = tempfile.mkdtemp(prefix="fx2trt.")
90+
self._data = ""
9091
_LOGGER.info(f"Initializing DiagnosticsWriter with root_dir: {self._root_dir}")
9192

9293
def write(self, file_name: str, data: WriteObj):
9394
"""
9495
TODO: Can be disabled by regex on file_name
9596
"""
97+
self._data = data
9698
# Only write if we are inside a collect_when() context.
9799
if not _IS_IN_COLLECT_CONTEXT.get(False):
98100
return
@@ -117,6 +119,9 @@ def write(self, file_name: str, data: WriteObj):
117119
def root_dir(self) -> str:
118120
return self._root_dir
119121

122+
def data(self) -> WriteObj:
123+
return self._data
124+
120125
def _write(self, file_name: str, to_write: bytes):
121126
# ms granularity - no naming collash, otherwise file will be
122127
# overwritten.
@@ -271,6 +276,9 @@ def collect(self) -> str:
271276
finally:
272277
os.remove(fp)
273278

279+
def data(self) -> WriteObj:
280+
return self._write.data()
281+
274282

275283
def _res_or_err(data: WriteObj) -> t.Tuple[TWrite, str]:
276284
if isinstance(data, (str, bytes)):

py/torch_tensorrt/fx/fx2trt.py

+6
Original file line numberDiff line numberDiff line change
@@ -1,4 +1,5 @@
11
import logging
2+
import os
23
import warnings
34
from datetime import datetime
45
from typing import Any, Callable, Dict, List, NamedTuple, Optional, Sequence
@@ -211,6 +212,11 @@ def run(
211212
builder_config = self.builder.create_builder_config()
212213
builder_config.max_workspace_size = max_workspace_size
213214

215+
# Speed up TRT build time in the test environment
216+
if trt.__version__ >= "8.6" and os.environ.get("TRT_TEST_ENV", "0") == "1":
217+
_LOGGER.info("Set TRT optimization level to 0")
218+
builder_config.builder_optimization_level = 0
219+
214220
cache = None
215221
if timing_cache:
216222
cache_file = numpy.array(timing_cache)

py/torch_tensorrt/fx/input_tensor_spec.py

+62-11
Original file line numberDiff line numberDiff line change
@@ -1,4 +1,4 @@
1-
from typing import Iterable, List, NamedTuple, Optional, Sequence, Tuple
1+
from typing import Any, Iterable, List, NamedTuple, Optional, Sequence, Tuple
22

33
import torch
44

@@ -18,6 +18,12 @@ def generate_input_specs(inputs, lower_setting, additional_inputs=None):
1818
# is the dynamic batch dimension. Otherwise, we use the additional
1919
# inputs to determine the batch dimension.
2020
if additional_inputs is None:
21+
batch_dims = None
22+
if not isinstance(inputs, torch.Tensor) and len(inputs) > 1:
23+
bs = inputs[0].size(0)
24+
batch_dims = None
25+
if not all(x.size(0) == bs for x in inputs):
26+
batch_dims = InputTensorSpec.find_batch_size_dim(inputs)
2127
return InputTensorSpec.from_tensors_with_dynamic_batch_size(
2228
inputs,
2329
(
@@ -26,6 +32,7 @@ def generate_input_specs(inputs, lower_setting, additional_inputs=None):
2632
lower_setting.max_batch_size,
2733
),
2834
lower_setting.opt_profile_replica,
35+
batch_dims,
2936
)
3037
else:
3138
batch_dims = []
@@ -147,25 +154,69 @@ def from_tensors_with_dynamic_batch_size(
147154
A list of InputTensorSpec named tuples with dynamic ranges.
148155
"""
149156
if batch_dims is None:
150-
batch_dims = [0] * len(tensors)
157+
batch_dims = cls.find_batch_size_dim(tensors)
151158

152159
input_specs = []
153160
batch_size = tensors[0].size(batch_dims[0])
154161

155162
for i, tensor in enumerate(tensors):
156163
batch_dim = batch_dims[i]
157-
assert batch_size == tensor.size(
158-
batch_dim
159-
), f"The {i}th tensor (shape: {tensor.shape}) doesn't have the correct batch size: {batch_size}."
160-
shape = list(tensor.shape)
161-
shape[batch_dim] = -1
162-
shape_ranges: List[ShapeRange] = [tuple(tuple(shape[0:batch_dim] + [bs] + shape[batch_dim + 1 :]) for bs in batch_size_range)] * opt_profile_replica # type: ignore[list-item]
163-
input_specs.append(
164-
cls(tuple(shape), tensor.dtype, tensor.device, shape_ranges)
165-
)
164+
if batch_dim == -1:
165+
input_specs.append(cls.from_tensor(tensor))
166+
else:
167+
shape = list(tensor.shape)
168+
assert batch_size == tensor.size(
169+
batch_dim
170+
), f"The {i}th tensor (shape: {tensor.shape}) doesn't have the correct batch size: {batch_size}."
171+
shape[batch_dim] = -1
172+
shape_ranges: List[ShapeRange] = [tuple(tuple(shape[0:batch_dim] + [bs] + shape[batch_dim + 1 :]) for bs in batch_size_range)] * opt_profile_replica # type: ignore[list-item]
173+
input_specs.append(
174+
cls(tuple(shape), tensor.dtype, tensor.device, shape_ranges)
175+
)
166176

167177
return input_specs
168178

179+
@classmethod
180+
# pyre-ignore [2]: Parameter `sample_input` must have a type other than `Any`
181+
def find_batch_size_dim(cls, inputs: Any) -> []:
182+
if isinstance(inputs, torch.Tensor) or len(inputs) <= 1:
183+
return [0]
184+
shapes = [i.shape for i in inputs]
185+
frequency_map = {}
186+
first_dims = set()
187+
for shape in shapes:
188+
if len(shape) < 2:
189+
# By pass for rank-1 tensors. MRS model has rank-1 tensor carry no batch_size info
190+
continue
191+
# Dedup shape value for single tensor
192+
first_dims.add(shape[0])
193+
shape = set(shape)
194+
for i in shape:
195+
frequency_map[i] = frequency_map.get(i, 0) + 1
196+
197+
if len(first_dims) == 1:
198+
# first dim is the same in every input: we use it as batch_size
199+
batch_size = first_dims.pop()
200+
elif frequency_map:
201+
# first dims are different: we use the most frequent dim as batch_size
202+
sorted_frequency = sorted(frequency_map.items(), key=lambda x: -x[1])
203+
batch_size = sorted_frequency[0][0]
204+
else:
205+
# no dims to sort: no batch_size
206+
batch_size = -1
207+
208+
bs_dim = []
209+
for i in inputs:
210+
# Default batch size dim = -1, indicate no batch_size
211+
dim = -1
212+
for index, val in enumerate(i.shape):
213+
if val == batch_size:
214+
dim = index
215+
break
216+
bs_dim.append(dim)
217+
218+
return bs_dim
219+
169220
def to_random_tensor(self, id=1):
170221
shape = tuple(self.shape)
171222
if len(get_dynamic_dims(shape)):

py/torch_tensorrt/fx/lower.py

+4
Original file line numberDiff line numberDiff line change
@@ -41,6 +41,8 @@ def compile(
4141
dynamic_batch=True,
4242
is_aten=False,
4343
use_experimental_fx_rt=False,
44+
correctness_atol=1e-1,
45+
correctness_rtol=1e-1,
4446
) -> nn.Module:
4547
"""
4648
Takes in original module, input and lowering setting, run lowering workflow to turn module
@@ -81,6 +83,8 @@ def compile(
8183
dynamic_batch=dynamic_batch,
8284
is_aten=is_aten,
8385
use_experimental_rt=use_experimental_fx_rt,
86+
correctness_atol=correctness_atol,
87+
correctness_rtol=correctness_rtol,
8488
)
8589
lowerer = Lowerer.create(lower_setting=lower_setting)
8690
return lowerer(module, input)

py/torch_tensorrt/fx/passes/lower_basic_pass.py

+41-5
Original file line numberDiff line numberDiff line change
@@ -54,10 +54,14 @@ def fill_with_mul_zero_and_add(*args):
5454

5555

5656
def run_const_fold(traced_mod: torch.fx.GraphModule) -> torch.fx.GraphModule:
57-
# Now we do constant folding on traced module. We want to skip pattern like
58-
# weights -> quant -> dequant -> op during constant folding when the model is
59-
# a quantized int8 model.
60-
def skip_folding_quant_dequant(node: torch.fx.Node):
57+
def skip_folding_ops(node: torch.fx.Node):
58+
# dtype op
59+
if node.target == acc_ops.dtype:
60+
return True
61+
# Now we do constant folding on traced module. We want to skip pattern like
62+
# weights -> quant -> dequant -> op during constant folding when the model is
63+
# a quantized int8 model.
64+
# quant_dequant
6165
if node.target != acc_ops.quantize_per_tensor:
6266
return False
6367
# If quantize_per_node -> dequantize, then skip folding.
@@ -66,7 +70,7 @@ def skip_folding_quant_dequant(node: torch.fx.Node):
6670
return True
6771
return False
6872

69-
const_split_mod = split_const_subgraphs(traced_mod, skip_folding_quant_dequant)
73+
const_split_mod = split_const_subgraphs(traced_mod, skip_folding_ops)
7074
const_split_mod.run_folding()
7175
return const_split_mod
7276

@@ -630,3 +634,35 @@ def fix_clamp_numerical_limits_to_fp16(
630634

631635
mod.recompile()
632636
return mod
637+
638+
639+
@log_before_after
640+
@validate_inference(atol=1e-3, rtol=1e-2)
641+
def remove_dtype_and_to_pattern(
642+
mod: torch.fx.GraphModule, input: Input
643+
) -> torch.fx.GraphModule:
644+
"""
645+
Remove this pattern since it is unnecessary to cast to dtype
646+
%dtype : [#users=1] = call_function[target=torch_tensorrt.fx.tracer.acc_tracer.acc_ops.dtype](args = (), kwargs = {input: %_attention_layers_0__uva})
647+
%to_18 : [#users=2] = call_function[target=torch_tensorrt.fx.tracer.acc_tracer.acc_ops.to_dtype](args = (), kwargs = {input: %x})
648+
"""
649+
for node in mod.graph.nodes:
650+
if node.op == "call_function" and node.target == acc_ops.dtype:
651+
# find its first user
652+
next_node = next(iter(node.users))
653+
# acc_op or pt op is treated differently
654+
input = (
655+
next_node.kwargs["input"]
656+
if "input" in next_node.kwargs
657+
else next_node.args[0]
658+
)
659+
if len(node.users) == 1 and (
660+
next_node.target == acc_ops.to_dtype or next_node.target == "to"
661+
):
662+
next_node.replace_all_uses_with(input)
663+
mod.graph.erase_node(next_node)
664+
mod.graph.erase_node(node)
665+
666+
mod.graph.eliminate_dead_code()
667+
mod.recompile()
668+
return mod

py/torch_tensorrt/fx/passes/lower_pass_manager_builder.py

+5-2
Original file line numberDiff line numberDiff line change
@@ -8,6 +8,7 @@
88
from torch.fx.passes.pass_manager import inplace_wrapper, PassManager
99
from torch.fx.passes.shape_prop import ShapeProp
1010
from torch.fx.passes.splitter_base import generate_inputs_for_submodules, SplitResult
11+
from torch_tensorrt.fx.passes.pass_utils import apply_bfloat_float_conversion
1112
from torch_tensorrt.fx.utils import LowerPrecision
1213

1314
from ..input_tensor_spec import generate_input_specs
@@ -229,10 +230,9 @@ def lower_func(split_result: SplitResult) -> nn.Module:
229230
submod = getattr(split_result.split_module, submod_name)
230231

231232
LOWER_SPLIT_PRE_OBSERVER.observe(submod_name, submod, submod_inputs)
232-
233233
# Only acc submodules will be lowered.
234234
if not submod_name.startswith(split_result.non_acc_submodule_prefix):
235-
_LOGGER.info(f"Now lowering submodule {submod_name}")
235+
_LOGGER.info(f"ACC submodule graph: {submod.graph}")
236236
lowering_start_time = datetime.datetime.now()
237237

238238
self.lower_setting.additional_inputs = (
@@ -251,6 +251,9 @@ def lower_func(split_result: SplitResult) -> nn.Module:
251251
_LOGGER.info(
252252
f"Lowering submodule {submod_name} elapsed time {datetime.datetime.now() - lowering_start_time}"
253253
)
254+
else:
255+
_LOGGER.info(f"GPU submodule graph: {submod.graph}")
256+
apply_bfloat_float_conversion(submod, submod_inputs, submod_name)
254257

255258
return split_result.split_module
256259

0 commit comments

Comments
 (0)