Skip to content

Commit f238716

Browse files
committed
add unbacked option
Signed-off-by: Laith Sakka <lsakka@meta.com>
1 parent 1017bb6 commit f238716

File tree

6 files changed

+344
-21
lines changed

6 files changed

+344
-21
lines changed
Lines changed: 145 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -0,0 +1,145 @@
1+
# SPDX-License-Identifier: Apache-2.0
2+
# SPDX-FileCopyrightText: Copyright contributors to the vLLM project
3+
4+
import gc
5+
6+
import pytest
7+
import torch
8+
from torch.torch_version import TorchVersion
9+
10+
from vllm import LLM, SamplingParams
11+
from vllm.config.compilation import DynamicShapesType
12+
13+
14+
def cleanup_gpu_memory():
15+
"""Clean up GPU memory after each test"""
16+
gc.collect() # Clear Python objects
17+
torch.cuda.empty_cache() # Clear PyTorch GPU memory cache
18+
torch.cuda.synchronize() # Wait for all GPU operations to complete
19+
20+
21+
def get_test_models():
22+
"""Get list of models to test based on PyTorch version"""
23+
# Parse PyTorch version
24+
result = ["microsoft/DialoGPT-small", "gpt2", "facebook/opt-125m"]
25+
# Handle alpha versions by removing pre-release suffixes
26+
version_parts = torch.__version__.split("+")[0].split("a")[0]
27+
clean_version = version_parts.split("b")[0].split("rc")[0]
28+
if TorchVersion(clean_version) >= TorchVersion("2.10"):
29+
# Requires some fixes only available in PyTorch 2.10+
30+
result.append("Qwen/Qwen2-1.5B-Instruct")
31+
result.append("Qwen/Qwen2-7B-Instruct")
32+
result.append("openlm-research/open_llama_13b")
33+
34+
return result
35+
36+
37+
@pytest.mark.parametrize("model_name", get_test_models())
38+
def test_dynamic_shapes_compilation(monkeypatch, model_name):
39+
"""Test that all dynamic shapes types produce compiles"""
40+
print(f"\nTesting model: {model_name}")
41+
42+
monkeypatch.setenv("TOKENIZERS_PARALLELISM", "true")
43+
# Note USE_AOT_COMPILE fails https://github.com/vllm-project/vllm/issues/27040.
44+
monkeypatch.setenv("VLLM_USE_AOT_COMPILE", "0")
45+
46+
prompt = "Hello, my name is"
47+
results = {}
48+
49+
print("Testing EAGER (no compilation) baseline...")
50+
cleanup_gpu_memory()
51+
52+
eager_model = LLM(
53+
model=model_name,
54+
compilation_config={
55+
"level": 0, # NO_COMPILATION - eager mode
56+
},
57+
# gpu_memory_utilization=0.2,
58+
)
59+
60+
# Generate text with deterministic sampling parameters
61+
sampling_params = SamplingParams(
62+
max_tokens=10,
63+
temperature=0.0, # Deterministic generation
64+
seed=42, # Fixed seed for consistency
65+
)
66+
eager_output = eager_model.generate(prompt, sampling_params=sampling_params)
67+
results["EAGER"] = eager_output[0].outputs[0].text
68+
69+
# Cleanup model
70+
del eager_model
71+
cleanup_gpu_memory()
72+
73+
# Test all dynamic shapes types with compilation
74+
for shapes_type in [
75+
DynamicShapesType.BACKED,
76+
DynamicShapesType.UNBACKED,
77+
DynamicShapesType.BACKED_SIZE_OBLIVIOUS,
78+
]:
79+
print(f"Testing {shapes_type.name} dynamic shapes...")
80+
81+
# Initialize the model with specific dynamic shapes configuration
82+
model = LLM(
83+
model=model_name,
84+
compilation_config={
85+
"level": 3, # PIECEWISE compilation
86+
"dynamic_shapes_config": {
87+
"dynamic_shapes_type": shapes_type.value,
88+
"eval_dynamo_ds_guards": False,
89+
},
90+
},
91+
# gpu_memory_utilization=0.2,
92+
)
93+
94+
output = model.generate(prompt, sampling_params=sampling_params)
95+
96+
# Store results for comparison
97+
results[shapes_type.name] = output[0].outputs[0].text
98+
99+
# Cleanup model
100+
del model
101+
cleanup_gpu_memory()
102+
103+
# Verify all results are non-empty strings
104+
for shape_type, result in results.items():
105+
assert isinstance(result, str), f"{shape_type} should return a string"
106+
assert len(result.strip()) > 0, f"{shape_type} should generate non-empty text"
107+
108+
# Print results
109+
for shape_type, result in results.items():
110+
print(f"{shape_type}: '{result}'")
111+
112+
113+
if __name__ == "__main__":
114+
"""Run the test directly as a Python script"""
115+
import os
116+
117+
print("Running dynamic shapes compilation test...")
118+
119+
# Get test models based on PyTorch version
120+
test_models = get_test_models()
121+
print(f"Testing {len(test_models)} models: {test_models}")
122+
123+
# Create a mock monkeypatch object for environment variables
124+
class MockMonkeypatch:
125+
def setenv(self, key, value):
126+
os.environ[key] = value
127+
128+
monkeypatch = MockMonkeypatch()
129+
130+
# Run test for each model
131+
for model_name in test_models:
132+
try:
133+
print(f"\n{'=' * 60}")
134+
print(f"Testing model: {model_name}")
135+
print(f"{'=' * 60}")
136+
137+
test_dynamic_shapes_compilation(monkeypatch, model_name)
138+
139+
print(f"✅ Test passed for {model_name}")
140+
141+
except Exception as e:
142+
print(f"❌ Test failed for {model_name}: {e}")
143+
raise
144+
145+
print("\n🎉 All tests completed successfully!")

vllm/compilation/decorators.py

Lines changed: 43 additions & 10 deletions
Original file line numberDiff line numberDiff line change
@@ -19,6 +19,7 @@
1919
from vllm.compilation.counter import compilation_counter
2020
from vllm.compilation.wrapper import TorchCompileGuardsStripWrapper
2121
from vllm.config import CompilationMode, VllmConfig, set_current_vllm_config
22+
from vllm.config.compilation import DynamicShapesType
2223
from vllm.logger import init_logger
2324
from vllm.sequence import IntermediateTensors
2425
from vllm.utils import resolve_obj_by_qualname, supports_dynamo
@@ -82,6 +83,7 @@ def support_torch_compile(
8283
*,
8384
dynamic_arg_dims: dict[str, int | list[int]] | None = None,
8485
enable_if: Callable[[VllmConfig], bool] | None = None,
86+
shape_invariants: Callable[..., None] = lambda *args, **kwargs: None,
8587
) -> Callable[[_T], _T] | _T:
8688
"""
8789
A decorator to add support for compiling the forward method of a class.
@@ -171,7 +173,9 @@ def cls_decorator_helper(cls: _T) -> _T:
171173
raise ValueError(
172174
f"Argument {k} not found in the forward method of {cls}"
173175
)
174-
return _support_torch_compile(cls, inferred_dynamic_arg_dims, enable_if)
176+
return _support_torch_compile(
177+
cls, inferred_dynamic_arg_dims, enable_if, shape_invariants
178+
)
175179

176180
if cls is not None:
177181
# use `support_torch_compile` as a decorator without arguments
@@ -212,6 +216,7 @@ def _support_torch_compile(
212216
cls: _T,
213217
dynamic_arg_dims: dict[str, int | list[int]],
214218
enable_if: Callable[[VllmConfig], bool] | None = None,
219+
shape_invariants: Callable[..., None] = lambda *args, **kwargs: None,
215220
) -> _T:
216221
"""
217222
A decorator to add support for compiling the forward method of a class.
@@ -232,11 +237,12 @@ def _support_torch_compile(
232237
def __init__(self, *, vllm_config: VllmConfig, prefix: str = "", **kwargs):
233238
old_init(self, vllm_config=vllm_config, prefix=prefix, **kwargs)
234239
self.vllm_config = vllm_config
240+
self.compilation_config = self.vllm_config.compilation_config
235241
enable_compile = enable_if is None or enable_if(vllm_config)
236242
# for CompilationMode.STOCK_TORCH_COMPILE , the upper level model runner
237243
# will handle the compilation, so we don't need to do anything here.
238244
self.do_not_compile = (
239-
vllm_config.compilation_config.mode
245+
self.compilation_config.mode
240246
in [CompilationMode.NONE, CompilationMode.STOCK_TORCH_COMPILE]
241247
or not supports_dynamo()
242248
or _should_ignore_torch_compile(self.__class__)
@@ -245,29 +251,38 @@ def __init__(self, *, vllm_config: VllmConfig, prefix: str = "", **kwargs):
245251
if self.do_not_compile:
246252
return
247253

254+
self._check_shape_invariants = shape_invariants
255+
248256
compilation_counter.num_models_seen += 1
249257
self.compiled = False
250258
TorchCompileGuardsStripWrapper.__init__(self)
251259

252260
cls.__init__ = __init__
253261

254-
def _mark_dynamic_inputs(mod, *args, **kwargs):
262+
def _mark_dynamic_inputs(mod, dynamic_shapes_type, *args, **kwargs):
263+
def mark_dynamic(arg, dims):
264+
if dynamic_shapes_type == DynamicShapesType.UNBACKED:
265+
torch._dynamo.decorators.mark_unbacked(arg, dims)
266+
else:
267+
torch._dynamo.mark_dynamic(arg, dims)
268+
255269
sig = inspect.signature(mod.__class__.forward)
256270
bound_args = sig.bind(mod, *args, **kwargs)
257271
bound_args.apply_defaults()
258272
for k, dims in dynamic_arg_dims.items():
259273
arg = bound_args.arguments.get(k)
274+
260275
if arg is not None:
261276
dims = [dims] if isinstance(dims, int) else dims
262277
if isinstance(arg, torch.Tensor):
263278
# In case dims is specified with negative indexing
264279
dims = [arg.ndim + dim if dim < 0 else dim for dim in dims]
265-
torch._dynamo.mark_dynamic(arg, dims)
280+
mark_dynamic(arg, dims)
266281
elif isinstance(arg, IntermediateTensors):
267282
for tensor in arg.tensors.values():
268283
# In case dims is specified with negative indexing
269284
dims = [tensor.ndim + dim if dim < 0 else dim for dim in dims]
270-
torch._dynamo.mark_dynamic(tensor, dims)
285+
mark_dynamic(tensor, dims)
271286
else:
272287
raise ValueError(
273288
"Unsupported dynamic dimensions"
@@ -285,6 +300,7 @@ def __call__(self, *args, **kwargs):
285300
if getattr(self, "aot_compiled_fn", None) is not None:
286301
return self.aot_compiled_fn(self, *args, **kwargs)
287302

303+
ds_type = self.compilation_config.dynamic_shapes_config.dynamic_shapes_type
288304
cache_dir = None
289305
aot_compilation_path = None
290306
if envs.VLLM_USE_AOT_COMPILE:
@@ -299,6 +315,14 @@ def __call__(self, *args, **kwargs):
299315
serialized backend artifacts), then we need to generate a new AOT
300316
compile artifact from scratch.
301317
"""
318+
# Validate that AOT compile is not used with unbacked dynamic
319+
# shapes. aot_compile re-allocates backed symbols post dynamo!
320+
if ds_type == DynamicShapesType.UNBACKED:
321+
raise ValueError(
322+
"AOT compilation is not compatible with UNBACKED dynamic shapes. "
323+
"Please use BACKED or BACKED_SIZE_OBLIVIOUS dynamic shapes type "
324+
"when VLLM_USE_AOT_COMPILE is enabled."
325+
)
302326
from .caching import compilation_config_hash_factors
303327

304328
factors: list[str] = compilation_config_hash_factors(self.vllm_config)
@@ -347,7 +371,12 @@ def __call__(self, *args, **kwargs):
347371
# This is the path for the first compilation.
348372

349373
# the first compilation needs to have dynamic shapes marked
350-
_mark_dynamic_inputs(self, *args, **kwargs)
374+
_mark_dynamic_inputs(
375+
self,
376+
ds_type,
377+
*args,
378+
**kwargs,
379+
)
351380

352381
# here, it is the starting point of the `torch.compile` process
353382
start_monitoring_torch_compile(self.vllm_config)
@@ -364,9 +393,7 @@ def __call__(self, *args, **kwargs):
364393
# properly when any of these files change.
365394

366395
# 1. the file containing the top-level forward function
367-
self.vllm_config.compilation_config.traced_files.add(
368-
original_code_object.co_filename
369-
)
396+
self.compilation_config.traced_files.add(original_code_object.co_filename)
370397

371398
# 2. every time Dynamo sees a function call, it will inline
372399
# the function by calling InliningInstructionTranslator.inline_call_
@@ -376,7 +403,7 @@ def __call__(self, *args, **kwargs):
376403

377404
def patched_inline_call(self_):
378405
code = self_.f_code
379-
self.vllm_config.compilation_config.traced_files.add(code.co_filename)
406+
self.compilation_config.traced_files.add(code.co_filename)
380407
return inline_call(self_)
381408

382409
# Disable the C++ compilation of symbolic shape guards. C++-fication
@@ -392,12 +419,18 @@ def patched_inline_call(self_):
392419
# if the config doesn't exist
393420
logger.debug("enable_cpp_symbolic_shape_guards config not available")
394421

422+
# Prepare backed_size_oblivious config patch if needed
423+
fx_config_patches = {}
424+
if ds_type == DynamicShapesType.BACKED_SIZE_OBLIVIOUS:
425+
fx_config_patches["backed_size_oblivious"] = True
426+
395427
with (
396428
patch.object(
397429
InliningInstructionTranslator, "inline_call_", patched_inline_call
398430
),
399431
torch._dynamo.config.patch(**dynamo_config_patches),
400432
maybe_use_cudagraph_partition_wrapper(self.vllm_config),
433+
torch.fx.experimental._config.patch(**fx_config_patches),
401434
_torch27_patch_tensor_subclasses(),
402435
):
403436
if envs.VLLM_USE_AOT_COMPILE:

vllm/compilation/wrapper.py

Lines changed: 7 additions & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -22,6 +22,12 @@ class TorchCompileGuardsStripWrapper:
2222
since we drop all guards.
2323
"""
2424

25+
def check_invariantes_and_forward(self, *args, **kwargs):
26+
assert hasattr(self, "_check_shape_invariants")
27+
self._check_shape_invariants(*args, **kwargs)
28+
29+
return self.forward(*args, **kwargs)
30+
2531
def __init__(self):
2632
self.compiled = False
2733

@@ -50,7 +56,7 @@ def __init__(self):
5056
logger.warning(msg)
5157

5258
self._compiled_callable = torch.compile(
53-
self.forward,
59+
self.check_invariantes_and_forward,
5460
fullgraph=True,
5561
backend=backend,
5662
options=options,

0 commit comments

Comments
 (0)