Skip to content

Commit 72947c5

Browse files
authored
chore: reenable py313 (#3455)
1 parent d3eb817 commit 72947c5

19 files changed

+245
-44
lines changed

.github/scripts/filter-matrix.py

+2-1
Original file line numberDiff line numberDiff line change
@@ -3,8 +3,9 @@
33
import argparse
44
import json
55
import sys
6+
from typing import List
67

7-
disabled_python_versions = "3.13"
8+
disabled_python_versions: List[str] = []
89

910

1011
def main(args: list[str]) -> None:

.github/scripts/generate-tensorrt-test-matrix.py

+8
Original file line numberDiff line numberDiff line change
@@ -28,6 +28,10 @@
2828
# please update the future tensorRT version you want to test here
2929
TENSORRT_VERSIONS_DICT = {
3030
"windows": {
31+
"10.3.0": {
32+
"urls": "https://developer.nvidia.com/downloads/compute/machine-learning/tensorrt/10.3.0/zip/TensorRT-10.3.0.26.Windows.win10.cuda-12.5.zip",
33+
"strip_prefix": "TensorRT-10.3.0.26",
34+
},
3135
"10.7.0": {
3236
"urls": "https://developer.nvidia.com/downloads/compute/machine-learning/tensorrt/10.7.0/zip/TensorRT-10.7.0.23.Windows.win10.cuda-12.6.zip",
3337
"strip_prefix": "TensorRT-10.7.0.23",
@@ -42,6 +46,10 @@
4246
},
4347
},
4448
"linux": {
49+
"10.3.0": {
50+
"urls": "https://developer.nvidia.com/downloads/compute/machine-learning/tensorrt/10.3.0/tars/TensorRT-10.3.0.26.Linux.x86_64-gnu.cuda-12.5.tar.gz",
51+
"strip_prefix": "TensorRT-10.3.0.26",
52+
},
4553
"10.7.0": {
4654
"urls": "https://developer.nvidia.com/downloads/compute/machine-learning/tensorrt/10.7.0/tars/TensorRT-10.7.0.23.Linux.x86_64-gnu.cuda-12.6.tar.gz",
4755
"strip_prefix": "TensorRT-10.7.0.23",

.github/scripts/generate_binary_build_matrix.py

+6-10
Original file line numberDiff line numberDiff line change
@@ -18,15 +18,16 @@
1818
import sys
1919
from typing import Any, Callable, Dict, List, Optional, Tuple
2020

21+
PYTHON_VERSIONS_FOR_PR_BUILD = ["3.11"]
2122
PYTHON_ARCHES_DICT = {
22-
"nightly": ["3.9", "3.10", "3.11", "3.12"],
23-
"test": ["3.9", "3.10", "3.11", "3.12"],
24-
"release": ["3.9", "3.10", "3.11", "3.12"],
23+
"nightly": ["3.9", "3.10", "3.11", "3.12", "3.13"],
24+
"test": ["3.9", "3.10", "3.11", "3.12", "3.13"],
25+
"release": ["3.9", "3.10", "3.11", "3.12", "3.13"],
2526
}
2627
CUDA_ARCHES_DICT = {
2728
"nightly": ["11.8", "12.6", "12.8"],
2829
"test": ["11.8", "12.6", "12.8"],
29-
"release": ["11.8", "12.6", "12.8"],
30+
"release": ["11.8", "12.4", "12.6"],
3031
}
3132
ROCM_ARCHES_DICT = {
3233
"nightly": ["6.1", "6.2"],
@@ -422,11 +423,6 @@ def generate_wheels_matrix(
422423
# Define default python version
423424
python_versions = list(PYTHON_ARCHES)
424425

425-
# If the list of python versions is set explicitly by the caller, stick with it instead
426-
# of trying to add more versions behind the scene
427-
if channel == NIGHTLY and (os in (LINUX, MACOS_ARM64, LINUX_AARCH64)):
428-
python_versions += ["3.13"]
429-
430426
if os == LINUX:
431427
# NOTE: We only build manywheel packages for linux
432428
package_type = "manywheel"
@@ -456,7 +452,7 @@ def generate_wheels_matrix(
456452
arches += [XPU]
457453

458454
if limit_pr_builds:
459-
python_versions = [python_versions[0]]
455+
python_versions = PYTHON_VERSIONS_FOR_PR_BUILD
460456

461457
global WHEEL_CONTAINER_IMAGES
462458

.github/workflows/build-test-linux.yml

-1
Original file line numberDiff line numberDiff line change
@@ -23,7 +23,6 @@ jobs:
2323
test-infra-ref: main
2424
with-rocm: false
2525
with-cpu: false
26-
python-versions: '["3.11", "3.12", "3.10", "3.9"]'
2726

2827
filter-matrix:
2928
needs: [generate-matrix]

.github/workflows/build-test-windows.yml

-1
Original file line numberDiff line numberDiff line change
@@ -23,7 +23,6 @@ jobs:
2323
test-infra-ref: main
2424
with-rocm: false
2525
with-cpu: false
26-
python-versions: '["3.11", "3.12", "3.10", "3.9"]'
2726

2827
substitute-runner:
2928
needs: generate-matrix

py/torch_tensorrt/_features.py

+19-1
Original file line numberDiff line numberDiff line change
@@ -14,6 +14,7 @@
1414
"torch_tensorrt_runtime",
1515
"dynamo_frontend",
1616
"fx_frontend",
17+
"refit",
1718
],
1819
)
1920

@@ -36,9 +37,10 @@
3637
_TORCHTRT_RT_AVAIL = _TS_FE_AVAIL or os.path.isfile(linked_file_runtime_full_path)
3738
_DYNAMO_FE_AVAIL = version.parse(sanitized_torch_version()) >= version.parse("2.1.dev")
3839
_FX_FE_AVAIL = True
40+
_REFIT_AVAIL = version.parse(sys.version.split()[0]) < version.parse("3.13")
3941

4042
ENABLED_FEATURES = FeatureSet(
41-
_TS_FE_AVAIL, _TORCHTRT_RT_AVAIL, _DYNAMO_FE_AVAIL, _FX_FE_AVAIL
43+
_TS_FE_AVAIL, _TORCHTRT_RT_AVAIL, _DYNAMO_FE_AVAIL, _FX_FE_AVAIL, _REFIT_AVAIL
4244
)
4345

4446

@@ -62,6 +64,22 @@ def not_implemented(*args: List[Any], **kwargs: Dict[str, Any]) -> Any:
6264
return wrapper
6365

6466

67+
def needs_refit(f: Callable[..., Any]) -> Callable[..., Any]:
68+
def wrapper(*args: List[Any], **kwargs: Dict[str, Any]) -> Any:
69+
if ENABLED_FEATURES.refit:
70+
return f(*args, **kwargs)
71+
else:
72+
73+
def not_implemented(*args: List[Any], **kwargs: Dict[str, Any]) -> Any:
74+
raise NotImplementedError(
75+
"Refit feature is currently not available in Python 3.13 or higher"
76+
)
77+
78+
return not_implemented(*args, **kwargs)
79+
80+
return wrapper
81+
82+
6583
T = TypeVar("T")
6684

6785

py/torch_tensorrt/dynamo/_refit.py

+12-3
Original file line numberDiff line numberDiff line change
@@ -11,6 +11,7 @@
1111
from torch.export import ExportedProgram
1212
from torch.fx.experimental.proxy_tensor import unset_fake_temporarily
1313
from torch_tensorrt._enums import dtype
14+
from torch_tensorrt._features import needs_refit
1415
from torch_tensorrt._Input import Input
1516
from torch_tensorrt.dynamo import partitioning
1617
from torch_tensorrt.dynamo._exporter import inline_torch_modules
@@ -47,6 +48,7 @@
4748
logger = logging.getLogger(__name__)
4849

4950

51+
@needs_refit
5052
def construct_refit_mapping(
5153
module: torch.fx.GraphModule,
5254
inputs: Sequence[Input],
@@ -108,8 +110,11 @@ def construct_refit_mapping(
108110
return weight_map
109111

110112

113+
@needs_refit
111114
def construct_refit_mapping_from_weight_name_map(
112-
weight_name_map: dict[Any, Any], state_dict: dict[Any, Any]
115+
weight_name_map: dict[Any, Any],
116+
state_dict: dict[Any, Any],
117+
settings: CompilationSettings,
113118
) -> dict[Any, Any]:
114119
engine_weight_map = {}
115120
for engine_weight_name, (sd_weight_name, np_weight_type) in weight_name_map.items():
@@ -120,7 +125,9 @@ def construct_refit_mapping_from_weight_name_map(
120125
# If weights is not in sd, we can leave it unchanged
121126
continue
122127
else:
123-
engine_weight_map[engine_weight_name] = state_dict[sd_weight_name]
128+
engine_weight_map[engine_weight_name] = state_dict[sd_weight_name].to(
129+
to_torch_device(settings.device)
130+
)
124131

125132
engine_weight_map[engine_weight_name] = (
126133
engine_weight_map[engine_weight_name]
@@ -134,6 +141,7 @@ def construct_refit_mapping_from_weight_name_map(
134141
return engine_weight_map
135142

136143

144+
@needs_refit
137145
def _refit_single_trt_engine_with_gm(
138146
new_gm: torch.fx.GraphModule,
139147
old_engine: trt.ICudaEngine,
@@ -163,7 +171,7 @@ def _refit_single_trt_engine_with_gm(
163171
"constant_mapping", {}
164172
) # type: ignore
165173
mapping = construct_refit_mapping_from_weight_name_map(
166-
weight_name_map, new_gm.state_dict()
174+
weight_name_map, new_gm.state_dict(), settings
167175
)
168176
constant_mapping_with_type = {}
169177

@@ -213,6 +221,7 @@ def _refit_single_trt_engine_with_gm(
213221
raise AssertionError("Refitting failed.")
214222

215223

224+
@needs_refit
216225
def refit_module_weights(
217226
compiled_module: torch.fx.GraphModule | ExportedProgram,
218227
new_weight_module: ExportedProgram,

py/torch_tensorrt/dynamo/conversion/_TRTInterpreter.py

+10-11
Original file line numberDiff line numberDiff line change
@@ -26,6 +26,7 @@
2626
from torch.fx.passes.shape_prop import TensorMetadata
2727
from torch.utils._python_dispatch import _disable_current_modes
2828
from torch_tensorrt._enums import dtype
29+
from torch_tensorrt._features import needs_refit
2930
from torch_tensorrt._Input import Input
3031
from torch_tensorrt.dynamo import _defaults
3132
from torch_tensorrt.dynamo._engine_cache import BaseEngineCache
@@ -44,7 +45,7 @@
4445
get_trt_tensor,
4546
to_torch,
4647
)
47-
from torch_tensorrt.dynamo.utils import DYNAMIC_DIM, get_model_device, to_torch_device
48+
from torch_tensorrt.dynamo.utils import DYNAMIC_DIM, to_torch_device
4849
from torch_tensorrt.fx.observer import Observer
4950
from torch_tensorrt.logging import TRT_LOGGER
5051

@@ -434,6 +435,7 @@ def check_weight_equal(
434435
except Exception:
435436
return torch.all(sd_weight == network_weight)
436437

438+
@needs_refit
437439
def _save_weight_mapping(self) -> None:
438440
"""
439441
Construct the weight name mapping from engine weight name to state_dict weight name.
@@ -491,15 +493,10 @@ def _save_weight_mapping(self) -> None:
491493
_LOGGER.info("Building weight name mapping...")
492494
# Stage 1: Name mapping
493495
torch_device = to_torch_device(self.compilation_settings.device)
494-
gm_is_on_cuda = get_model_device(self.module).type == "cuda"
495-
if not gm_is_on_cuda:
496-
# If the model original position is on CPU, move it GPU
497-
sd = {
498-
k: v.reshape(-1).to(torch_device)
499-
for k, v in self.module.state_dict().items()
500-
}
501-
else:
502-
sd = {k: v.reshape(-1) for k, v in self.module.state_dict().items()}
496+
sd = {
497+
k: v.reshape(-1).to(torch_device)
498+
for k, v in self.module.state_dict().items()
499+
}
503500
weight_name_map: dict[str, Any] = {}
504501
np_map = {}
505502
constant_mapping = {}
@@ -583,6 +580,7 @@ def _save_weight_mapping(self) -> None:
583580
gc.collect()
584581
torch.cuda.empty_cache()
585582

583+
@needs_refit
586584
def _insert_engine_to_cache(self, hash_val: str, serialized_engine: bytes) -> None:
587585
# TODO: @Evan is waiting for TRT's feature to cache the weight-stripped engine
588586
# if not self.compilation_settings.strip_engine_weights:
@@ -610,6 +608,7 @@ def _insert_engine_to_cache(self, hash_val: str, serialized_engine: bytes) -> No
610608
),
611609
)
612610

611+
@needs_refit
613612
def _pull_cached_engine(self, hash_val: str) -> Optional[TRTInterpreterResult]:
614613
# query the cached TRT engine
615614
cached_data = self.engine_cache.check(hash_val) # type: ignore[union-attr]
@@ -720,7 +719,7 @@ def run(
720719
if self.compilation_settings.reuse_cached_engines:
721720
interpreter_result = self._pull_cached_engine(hash_val)
722721
if interpreter_result is not None: # hit the cache
723-
return interpreter_result
722+
return interpreter_result # type: ignore[no-any-return]
724723

725724
self._construct_trt_network_def()
726725

pyproject.toml

+1-1
Original file line numberDiff line numberDiff line change
@@ -1,6 +1,6 @@
11
[build-system]
22
requires = [
3-
"setuptools>=68.0.0",
3+
"setuptools>=77.0.0",
44
"packaging>=23.1",
55
"wheel>=0.40.0",
66
"ninja>=1.11.0",

setup.py

+1-1
Original file line numberDiff line numberDiff line change
@@ -18,12 +18,12 @@
1818
import torch
1919
import yaml
2020
from setuptools import Extension, find_namespace_packages, setup
21+
from setuptools.command.bdist_wheel import bdist_wheel
2122
from setuptools.command.build_ext import build_ext
2223
from setuptools.command.develop import develop
2324
from setuptools.command.editable_wheel import editable_wheel
2425
from setuptools.command.install import install
2526
from torch.utils.cpp_extension import IS_WINDOWS, BuildExtension, CUDAExtension
26-
from wheel.bdist_wheel import bdist_wheel
2727

2828
__version__: str = "0.0.0"
2929
__cuda_version__: str = "0.0"

tests/modules/custom_models.py

+2-1
Original file line numberDiff line numberDiff line change
@@ -3,7 +3,6 @@
33
import torch
44
import torch.nn as nn
55
import torch.nn.functional as F
6-
from transformers import BertConfig, BertModel, BertTokenizer
76

87

98
# Sample Pool Model (for testing plugin serialization)
@@ -165,6 +164,8 @@ def forward(self, z: List[torch.Tensor]):
165164

166165

167166
def BertModule():
167+
from transformers import BertConfig, BertModel, BertTokenizer
168+
168169
enc = BertTokenizer.from_pretrained("google-bert/bert-base-uncased")
169170
text = "[CLS] Who was Jim Henson ? [SEP] Jim Henson was a puppeteer [SEP]"
170171
tokenized_text = enc.tokenize(text)

tests/modules/hub.py

-3
Original file line numberDiff line numberDiff line change
@@ -4,10 +4,7 @@
44
import custom_models as cm
55
import timm
66
import torch
7-
import torch.nn as nn
8-
import torch.nn.functional as F
97
import torchvision.models as models
10-
from transformers import BertConfig, BertModel, BertTokenizer
118

129
torch.hub._validate_not_a_forked_repo = lambda a, b, c: True
1310

tests/py/dynamo/models/test_engine_cache.py

+20
Original file line numberDiff line numberDiff line change
@@ -250,6 +250,10 @@ def remove_timing_cache(path=TIMING_CACHE_PATH):
250250
msg=f"Engine caching didn't speed up the compilation. Time taken without engine caching: {times[0]} ms, time taken with engine caching: {times[2]} ms",
251251
)
252252

253+
@unittest.skipIf(
254+
not torch_trt.ENABLED_FEATURES.refit,
255+
"Engine caching requires refit feature that is not supported in Python 3.13 or higher",
256+
)
253257
def test_dynamo_compile_with_custom_engine_cache(self):
254258
model = models.resnet18(pretrained=True).eval().to("cuda")
255259

@@ -314,6 +318,10 @@ def test_dynamo_compile_with_custom_engine_cache(self):
314318
for h, count in custom_engine_cache.hashes.items()
315319
]
316320

321+
@unittest.skipIf(
322+
not torch_trt.ENABLED_FEATURES.refit,
323+
"Engine caching requires refit feature that is not supported in Python 3.13 or higher",
324+
)
317325
def test_dynamo_compile_change_input_shape(self):
318326
"""Runs compilation 3 times, the cache should miss each time"""
319327
model = models.resnet18(pretrained=True).eval().to("cuda")
@@ -346,6 +354,10 @@ def test_dynamo_compile_change_input_shape(self):
346354
for h, count in custom_engine_cache.hashes.items()
347355
]
348356

357+
@unittest.skipIf(
358+
not torch_trt.ENABLED_FEATURES.refit,
359+
"Engine caching requires refit feature that is not supported in Python 3.13 or higher",
360+
)
349361
@pytest.mark.xfail
350362
def test_torch_compile_with_default_disk_engine_cache(self):
351363
# Custom Engine Cache
@@ -485,6 +497,10 @@ def test_torch_compile_with_custom_engine_cache(self):
485497
for h, count in custom_engine_cache.hashes.items()
486498
]
487499

500+
@unittest.skipIf(
501+
not torch_trt.ENABLED_FEATURES.refit,
502+
"Engine caching requires refit feature that is not supported in Python 3.13 or higher",
503+
)
488504
def test_torch_trt_compile_change_input_shape(self):
489505
# Custom Engine Cache
490506
model = models.resnet18(pretrained=True).eval().to("cuda")
@@ -611,6 +627,10 @@ def forward(self, c, d):
611627
assertions.assertEqual(hash1, hash2)
612628

613629
# @unittest.skip("benchmark on small models")
630+
@unittest.skipIf(
631+
not torch_trt.ENABLED_FEATURES.refit,
632+
"Engine caching requires refit feature that is not supported in Python 3.13 or higher",
633+
)
614634
def test_caching_small_model(self):
615635
from torch_tensorrt.dynamo._refit import refit_module_weights
616636

0 commit comments

Comments
 (0)