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 416 417 418 419 420 421 422 423 424 425 426 427 428 429 430 431 432 433 434 435 436 437 438 439 440 441 442 443 444 445 446 447 448 449 450 451 452 453 454 455 456 457 458 459 460 461 462 463 464 465 466 467 468 469 470 471 472 473 474 475 476 477 478 479 480 481 482 483 484 485 486 487 488 489 490 491 492 493 494 495 496 497 498 499 500 501 502 503 504 505 506 507 508 509 510 511 512 513 514 515 516 517 518 519 520 521 522 523 524 525 526 527 528 529 530 531 532 533 534 535 536 537 538 539 540 541 542 543 544 545 546 547 548 549 550 551 552 553 554 555 556 557 558 559 560 561 562 563 564 565 566 567 568 569 570 571 572 573 574 575 576 577 578 579 580 581 582 583 584 585 586 587 588 589 590 591 592 593 594 595 596 597 598 599 600 601 602 603 604 605 606 607 608 609 610 611 612 613 614 615 616 617 618 619 620 621 622 623 624 625 626 627 628 629 630 631 632 633 634 635 636 637 638 639 640 641 642 643 644 645 646 647 648 649 650 651 652 653 654 655 656 657 658 659 660 661 662 663 664 665 666 667 668 669 670 671 672 673 674 675 676 677 678 679 680 681 682 683 684 685 686 687 688 689 690 691 692 693 694 695 696 697 698 699 700 701 702 703 704 705 706 707 708 709 710 711 712 713 714 715 716 717 718 719 720 721 722 723 724 725 726 727 728 729 730 731 732 733 734 735 736 737 738 739 740 741 742 743 744 745 746 747 748 749 750 751 752 753 754 755 756 757 758 759 760 761 762 763 764 765 766 767 768 769 770 771 772 773 774 775 776 777 778 779 780 781 782 783 784 785 786 787 788 789 790 791 792 793 794 795 796 797 798 799 800 801 802 803 804 805 806 807 808 809 810 811 812 813 814 815 816 817 818 819 820 821 822 823 824 825 826 827
|
# Owner(s): ["module: dynamo"]
import functools
import operator
import os
import unittest.mock as mock
from unittest.mock import patch
import torch
import torch._dynamo.test_case
import torch._dynamo.testing
from torch._dynamo.exc import IncorrectUsage
from torch._dynamo.utils import counters
def my_custom_function(x):
return x + 1
class DecoratorTests(torch._dynamo.test_case.TestCase):
def test_disallow_in_graph(self):
cnts = torch._dynamo.testing.CompileCounter()
@torch.compile(backend=cnts)
def fn(a):
x = torch.add(a, 1)
x = torch.add(x, 1)
x = torch.sub(x, 1)
x = torch.add(x, 1)
x = torch.add(x, 1)
return x
torch._dynamo.disallow_in_graph(torch.sub)
fn(torch.randn(10))
torch._dynamo.allow_in_graph(torch.sub)
# check for graph break on sub
self.assertEqual(cnts.frame_count, 2)
self.assertEqual(cnts.op_count, 4)
def test_disable_for_custom_op(self):
import torch.library
from torch.library import Library
foo = Library("foo", "DEF") # noqa: TOR901
foo.define("custom(Tensor self) -> Tensor")
# Dynamic shape data dependent operator. For static shape compilation, Dynamo
# should graph break on it. But, the meta kernel is not implemented properly.
@torch.library.impl(foo, "custom", "CPU")
def foo_cpu(x):
return x.nonzero()
# Disallow does not work because of extra python frames with torch.library python API
torch.ops.foo.custom = torch._dynamo.disable(torch.ops.foo.custom)
def fn(x):
a = torch.nn.functional.relu(x)
b = torch.ops.foo.custom(a)
c = torch.cos(b)
return c
x = torch.randint(2, (100,))
ref = fn(x)
cnts = torch._dynamo.testing.CompileCounter()
opt_fn = torch.compile(fn, backend=cnts)
res = opt_fn(x)
self.assertEqual(cnts.frame_count, 2)
self.assertEqual(ref, res)
def test_disable_ignores_outer_wraps(self):
def orig_inner():
pass
def inner():
pass
inner._torchdynamo_orig_callable = orig_inner
@functools.wraps(inner)
def wrapper():
raise AssertionError("wrapper called")
# This behavior is not ideal, but supporting it would add overhead
# to callsites of eval_frame.innermost_fn. A warning would also be very noisy.
w = torch._dynamo.disable(fn=wrapper, recursive=True)
def test_disable_nn_modules_forward_hook(self):
class SimpleLinear(torch.nn.Module):
def __init__(self) -> None:
super().__init__()
self.layer0 = torch.nn.Linear(4, 4)
def forward(self, inp):
return self.layer0(torch.sigmoid(inp))
class SimpleModel(torch.nn.Module):
def __init__(self) -> None:
super().__init__()
self.layer0 = SimpleLinear()
self.layer1 = torch.nn.Linear(4, 4)
def forward(self, inp):
z = self.layer0(torch.sin(inp))
return self.layer1(z)
def hook(module, args):
inp = args[0].sigmoid()
return (inp,)
model = SimpleModel()
model.layer0.register_forward_pre_hook(hook)
# Disable my monkeypatching
model.layer0 = torch._dynamo.disable(model.layer0)
cnts = torch._dynamo.testing.CompileCounterWithBackend("eager")
opt_model = torch.compile(model, backend=cnts)
opt_model(torch.randn(4))
# check for no graph break
self.assertEqual(cnts.frame_count, 2)
gm0 = cnts.graphs[0]
# Check that the first graph has sin node, and no sigmoid
self.assertTrue(any(node.target is torch.sin for node in gm0.graph.nodes))
self.assertTrue(
all(node.target is not torch.sigmoid for node in gm0.graph.nodes)
)
gm1 = cnts.graphs[1]
# Check that the first graph does not have sigmoid. sigmoid is used in
# both hook and disabled module.
self.assertTrue(
all(node.target is not torch.sigmoid for node in gm1.graph.nodes)
)
def test_disable_nn_module_with_class_decorator(self):
cnts = torch._dynamo.testing.CompileCounterWithBackend("eager")
@torch._dynamo.disable
class SimpleLinear(torch.nn.Module):
def __init__(self) -> None:
super().__init__()
self.layer0 = torch.nn.Linear(4, 4)
def forward(self, inp):
return self.layer0(torch.sigmoid(inp))
@torch.compile(backend=cnts)
class SimpleModel(torch.nn.Module):
def __init__(self) -> None:
super().__init__()
self.layer0 = SimpleLinear()
self.layer1 = torch.nn.Linear(4, 4)
def forward(self, inp):
z = self.layer0(torch.sin(inp))
return self.layer1(z)
def hook(module, args):
inp = args[0].sigmoid()
return (inp,)
model = SimpleModel()
model.layer0.register_forward_pre_hook(hook)
model(torch.randn(4))
# check for no graph break
self.assertEqual(cnts.frame_count, 2)
gm0 = cnts.graphs[0]
# Check that the first graph has sin node, and no sigmoid
self.assertTrue(any(node.target is torch.sin for node in gm0.graph.nodes))
self.assertTrue(
all(node.target is not torch.sigmoid for node in gm0.graph.nodes)
)
gm1 = cnts.graphs[1]
# Check that the first graph does not have sigmoid. sigmoid is used in
# both hook and disabled module.
self.assertTrue(
all(node.target is not torch.sigmoid for node in gm1.graph.nodes)
)
def test_allow_in_graph(self):
cnts = torch._dynamo.testing.CompileCounter()
@torch.compile(backend=cnts)
def fn(a):
x = torch.add(a, 1)
x = torch.add(x, 1)
x = my_custom_function(x)
x = torch.add(x, 1)
x = torch.add(x, 1)
return x
torch._dynamo.allow_in_graph(my_custom_function)
fn(torch.randn(10))
torch._dynamo.disallow_in_graph(my_custom_function)
# check for no graph break
self.assertEqual(cnts.frame_count, 1)
self.assertEqual(cnts.op_count, 5)
def test_incorrect_usage_disallow_in_graph(self):
with self.assertRaises(IncorrectUsage):
@torch._dynamo.disallow_in_graph
def fn1(x):
return x.cos()
def test_graph_break(self):
cnts = torch._dynamo.testing.CompileCounter()
@torch.compile(backend=cnts)
def fn(x):
x = torch.cos(x)
x = torch.cos(x)
torch._dynamo.graph_break()
x = torch.cos(x)
x = torch.cos(x)
torch._dynamo.graph_break()
x = torch.cos(x)
x = torch.cos(x)
return x
fn(torch.randn(4, 5))
self.assertEqual(cnts.frame_count, 3)
self.assertEqual(cnts.op_count, 6)
def test_skip(self):
def fn2(x):
return x.sin()
@torch._dynamo.disable(recursive=False)
def fn1(x):
x = x.sigmoid()
return fn2(x.cos())
def fn(x):
return fn1(x.tan())
cnts = torch._dynamo.testing.CompileCounter()
opt_fn = torch.compile(fn, backend=cnts)
opt_fn(torch.randn(4))
self.assertEqual(cnts.frame_count, 2)
def test_substitute_in_graph(self):
counters.clear()
# NB: Choose another C function for test when we support operator.indexOf
# out of the box
cnts = torch._dynamo.testing.CompileCounter()
fn = operator.indexOf
opt_fn = torch.compile(fn, backend=cnts)
out = fn([1, 2, 3, 4, 5], 3)
opt_out = opt_fn([1, 2, 3, 4, 5], 3)
self.assertEqual(out, opt_out)
self.assertEqual(cnts.frame_count, 0)
self.assertEqual(len(counters["graph_break"]), 1)
torch._dynamo.reset()
counters.clear()
with self.assertRaisesRegex(TypeError, "Signature mismatch"):
@torch._dynamo.substitute_in_graph(operator.indexOf)
def _(sequence, x):
for i, item in enumerate(sequence):
if item is x or item == x:
return i
raise ValueError("sequence.index(x): x not in sequence")
@torch._dynamo.substitute_in_graph(operator.indexOf)
def polyfill(a, b):
for i, item in enumerate(a):
if item is b or item == b:
return i
raise ValueError("sequence.index(x): x not in sequence")
cnts = torch._dynamo.testing.CompileCounter()
fn = operator.indexOf
opt_fn = torch.compile(fn, backend=cnts, fullgraph=True)
out = fn([1, 2, 3, 4, 5], 3)
opt_out = opt_fn([1, 2, 3, 4, 5], 3)
self.assertEqual(out, opt_out)
self.assertEqual(cnts.frame_count, 0)
self.assertEqual(len(counters["graph_break"]), 0)
torch._dynamo.reset()
counters.clear()
cnts = torch._dynamo.testing.CompileCounter()
fn = polyfill
opt_fn = torch.compile(fn, backend=cnts, fullgraph=True)
out = fn([1, 2, 3, 4, 5], 3)
opt_out = opt_fn([1, 2, 3, 4, 5], 3)
self.assertEqual(out, opt_out)
self.assertEqual(cnts.frame_count, 0)
self.assertEqual(len(counters["graph_break"]), 0)
@patch.object(torch._dynamo.config, "suppress_errors", True)
def test_nested_disable_decorator(self):
cnts = torch._dynamo.testing.CompileCounter()
@torch._dynamo.disable()
def fn1(x):
return torch.sin(x) * 10
@torch.compile(backend=cnts)
def fn2(x):
x = x + 1
x = x + 1
x = fn1(x) # graph break
x = x + 1
x = x + 1
return x
@torch.compile(backend=cnts, fullgraph=True)
def fn3(x):
return fn2(x)
fn2(torch.randn(4, 5))
self.assertEqual(cnts.frame_count, 2)
self.assertEqual(cnts.op_count, 4)
try:
fn3(torch.randn(4, 5))
self.assertFalse(True)
except torch._dynamo.exc.Unsupported as e:
self.assertIn("call torch._dynamo.disable() wrapped function", str(e))
def test_disable_optimize(self):
cnt = torch._dynamo.testing.CompileCounter()
@torch.compile(backend=cnt, disable=True)
def f1(x):
return x + 1
f1(torch.ones(6))
self.assertEqual(cnt.frame_count, 0)
@torch.compile(backend=cnt, disable=True)
def f2(x):
return x + 1
f2(torch.ones(6))
self.assertEqual(cnt.frame_count, 0)
with patch.dict(os.environ, {"TORCHDYNAMO_DISABLE": "1"}):
@torch.compile(backend=cnt)
def f3(x):
return x + 1
f3(torch.ones(6))
self.assertEqual(cnt.frame_count, 0)
def test_torch_guards_stack_frame_register_inlining_disable(self):
x = torch.tensor([0.5, 0.5])
class encoder(torch.nn.Module):
def __init__(self, y):
super().__init__()
self.a = y
@torch._dynamo.disable
def helper(self, x, y):
return x * y
def forward(self, a, *args):
x = a + a
return self.helper(x, self.a)
e = encoder(2.0)
seen_frames = []
import contextlib
@contextlib.contextmanager
def global_context_capture_fn(frame_summary):
if frame_summary is not None:
seen_frames.append(frame_summary)
yield
with mock.patch(
"torch._guards.TracingContext.current_frame",
side_effect=global_context_capture_fn,
):
torch.compile(e, backend="eager")(x)
self.assertEqual(len(seen_frames), 0)
def test_torch_guards_stack_frame_register_inlining_partially_disable(self):
y = torch.nn.Parameter(torch.tensor([0.25, 0.25]))
x = torch.tensor([0.5, 0.5])
class encoder(torch.nn.Module):
def __init__(self, y):
super().__init__()
self.register_parameter("param", y)
@torch._dynamo.disable
def helper_disabled(self, x, y):
return x.sin() * y.cos()
def helper(self, x, y):
return x * y
def forward(self, a, *args):
x = a + a
return self.helper(x, self.param) + self.helper_disabled(x, self.param)
e = encoder(y)
cnt = torch._dynamo.testing.CompileCounter()
torch.compile(e, backend=cnt)(x)
# first frame is before disable, second frame is after disable
self.assertEqual(cnt.frame_count, 2)
self.assertEqual(cnt.op_count, 3)
def _test_mark_static_address(self, guarded):
# This test verifies that dynamo properly marks inputs as static
# when using the mark_static_address API.
# On 1st compile, we expect the input to be marked as static, with guarded
# set depending on the `guarded` flag.
# On 2nd compile, we expect the input to be unmarked
# if inlining NN modules, we expect metadata to be present on the tensor, indicating
# the static address type of the input
# if not inlining NN modules, we expect the tensor to be present in the buffers attribute
# of the graph.
compiles_with_buffers = 0
compiles = 0
def debug_compiler(gm, _):
nonlocal compiles_with_buffers
nonlocal compiles
if torch._dynamo.config.inline_inbuilt_nn_modules:
input_node = [
n
for n in gm.graph.nodes
if n.op == "placeholder" and n.name == "l_x_"
]
self.assertEqual(len(input_node), 1)
input_node = input_node[0]
if compiles == 0:
self.assertEqual(
input_node.meta["tensor_dict"]["_dynamo_static_input_type"],
"guarded" if guarded else "unguarded",
)
elif compiles == 1:
self.assertFalse(
"_dynamo_static_input_type" in input_node.meta["tensor_dict"]
)
else:
raise RuntimeError(f"Unexpected number of compiles: {compiles}")
else:
compiles_with_buffers += len(gm._buffers) > 0
compiles += 1
return gm
@torch.compile(backend=debug_compiler)
def fn(x):
return x + 1
inp = torch.ones(2)
torch._dynamo.mark_static_address(inp, guard=guarded)
fn(inp)
if not torch._dynamo.config.inline_inbuilt_nn_modules:
self.assertEqual(compiles_with_buffers, 1)
inp2 = torch.ones(2)
# if guarded, should trigger another recompile
# since it was not marked static, compiles with buffers
# should not be incremented
fn(inp2)
if not torch._dynamo.config.inline_inbuilt_nn_modules:
self.assertEqual(compiles_with_buffers, 1)
self.assertEqual(compiles, 2 if guarded else 1)
def test_mark_static_address_guarded(self):
with torch._dynamo.config.patch("inline_inbuilt_nn_modules", True):
self._test_mark_static_address(guarded=True)
self._test_mark_static_address(guarded=True)
def test_mark_static_address_unguarded(self):
with torch._dynamo.config.patch("inline_inbuilt_nn_modules", True):
self._test_mark_static_address(guarded=False)
self._test_mark_static_address(guarded=False)
def test_class_methods(self):
class A:
@classmethod
def my_class_method(cls, arg1):
return cls, arg1
@staticmethod
def my_static_method(arg1):
return None, arg1
def my_regular_method(self, arg1):
return self, arg1
class B(A):
def my_class_method(self, arg1):
return super().my_class_method(arg1)
def my_static_method(self, arg1):
return super().my_static_method(arg1)
class C(A):
@classmethod
def my_class_method(cls, arg1):
return super().my_class_method(arg1)
cnt = torch._dynamo.testing.CompileCounter()
@torch.compile(backend=cnt)
def fn(a, b, c):
# We want a function that does not graph break but
# does generate custom bytecode
v1 = a.my_class_method(1)
v2 = A.my_class_method(2)
v3 = a.my_static_method(3)
v4 = A.my_static_method(4)
v5 = a.my_regular_method(5)
v6 = b.my_class_method(6)
v7 = b.my_static_method(7)
v8 = c.my_class_method(8)
v9 = C.my_class_method(9)
torch.rand(2)
return v1, v2, v3, v4, v5, v6, v7, v8, v9
a, b, c = A(), B(), C()
v1, v2, v3, v4, v5, v6, v7, v8, v9 = fn(a, b, c)
self.assertEqual(v1, (A, 1))
self.assertEqual(v2, (A, 2))
self.assertEqual(v3, (None, 3))
self.assertEqual(v4, (None, 4))
self.assertEqual(v5, (a, 5))
# TODO fix me: we do not resolve classmethods properly
# from a regular method
# self.assertEqual(v6, (B, 6))
self.assertEqual(v7, (None, 7))
self.assertEqual(v8, (C, 8))
self.assertEqual(v9, (C, 9))
self.assertEqual(cnt.frame_count, 1)
def test_assume_constant_result_on_user_defined_fn(self):
@torch._dynamo.assume_constant_result
def const_fn(n, s):
return torch.full([n], s)
def fn(B):
B = const_fn(B.size(0), 13)
X = B * 2
return X.tolist()
B_list = [8] * 32
B = torch.tensor(B_list, dtype=torch.int32)
torch._dynamo.decorators.mark_static(B, 0)
torch._dynamo.config.capture_scalar_outputs = True
torch._dynamo.config.capture_dynamic_output_shape_ops = True
self.assertEqual(
fn(B), torch.compile(fn, backend="eager", fullgraph=True, dynamic=True)(B)
)
def test_assume_constant_result_on_computation_with_graph_input(self):
@torch._dynamo.assume_constant_result
def check(y):
return y[0].item() == 1
def fn(x, y):
if check(y):
return x + 2
else:
return x + 1
y = torch.tensor([1])
x = torch.tensor(1)
self.assertEqual(fn(x, y), torch.compile(fn)(x, y))
@torch._dynamo.config.patch("inline_inbuilt_nn_modules", True)
def test_mark_static_nn_module(self):
@torch._dynamo.mark_static
class Mock(torch.nn.Module):
def __init__(self, c):
super().__init__()
self.c = c
def forward(self, x):
return x * self.c
cnts = torch._dynamo.testing.CompileCounter()
mod1 = Mock(10)
mod2 = Mock(20)
mod3 = Mock(30)
opt_mod1 = torch.compile(mod1, backend=cnts, fullgraph=True)
opt_mod2 = torch.compile(mod2, backend=cnts, fullgraph=True)
opt_mod3 = torch.compile(mod3, backend=cnts, fullgraph=True)
x = torch.randn(4, 4)
opt_mod1(x)
opt_mod2(x)
opt_mod3(x)
# Must be 3 compilations. If not marked static there would be 2, because self.c would be converted to symints.
self.assertEqual(cnts.frame_count, 3)
def test_set_stance_force_eager(self):
@torch.compile(backend="eager")
def a(x):
if torch._dynamo.is_compiling():
return x + 1
return x + 2
@torch.compiler.set_stance("force_eager")
def b(x):
return a(x)
def c(x):
out0 = a(x)
with torch.compiler.set_stance("force_eager"):
out1 = a(x)
return out0, out1, a(x)
inp = torch.ones(3)
# test that decorating b has no overall side effect
self.assertEqual(a(inp), inp + 1)
self.assertEqual(b(inp), inp + 2)
self.assertEqual(c(inp), (inp + 1, inp + 2, inp + 1))
torch.compiler.set_stance("force_eager")
self.assertEqual(a(inp), inp + 2)
torch.compiler.set_stance("default")
self.assertEqual(a(inp), inp + 1)
def test_set_stance_eager_on_recompile(self):
@torch.compile(backend="eager", dynamic=False)
def a(x, n):
if torch._dynamo.is_compiling():
return x + n + 1
return x + n + 2
inp = torch.ones(3)
out1 = a(inp, 1)
with torch.compiler.set_stance("eager_on_recompile"):
out2 = a(inp, 1)
out3 = a(inp, 2)
self.assertEqual(out1, inp + 2)
self.assertEqual(out2, inp + 2)
self.assertEqual(out3, inp + 4)
def test_set_stance_fail_on_recompile(self):
@torch.compile(backend="eager", dynamic=False)
def a(x, n):
if torch._dynamo.is_compiling():
return x + n + 1
return x + n + 2
inp = torch.ones(3)
out1 = a(inp, 1)
with torch.compiler.set_stance("fail_on_recompile"):
out2 = a(inp, 1)
with self.assertRaisesRegex(RuntimeError, "fail_on_recompile"):
a(inp, 2)
self.assertEqual(out1, inp + 2)
self.assertEqual(out2, inp + 2)
def test_set_stance_fail_on_recompile_with_disable(self):
@torch.compiler.disable
def inner(x):
return x
@torch.compile(backend="eager")
def f(x):
return inner(x)
f(torch.randn(3, 3))
# should not raise error
with torch.compiler.set_stance("fail_on_recompile"):
f(torch.randn(3, 3))
def test_set_stance_forbid_in_graph(self):
@torch.compiler.set_stance("force_eager")
def a(x):
return x + 1
@torch.compile(backend="eager")
def b(x):
return a(x)
with self.assertRaisesRegex(
AssertionError, "Attempt to trace forbidden callable"
):
b(torch.ones(3))
@torch.compile(backend="eager")
def c(x):
with torch.compiler.set_stance("force_eager"):
return x + 1
with self.assertRaisesRegex(
AssertionError, "Attempt to trace forbidden callable"
):
c(torch.ones(3))
@torch.compile(backend="eager")
@torch.compiler.set_stance("force_eager")
def d(x):
return x + 1
with self.assertRaisesRegex(
AssertionError, "Attempt to trace forbidden callable"
):
d(torch.ones(3))
@torch.compile(backend="eager")
def e(x):
with torch._dynamo.set_stance("force_eager"):
return x + 1
with self.assertRaisesRegex(
AssertionError, "Attempt to trace forbidden callable"
):
e(torch.ones(3))
@torch.compile(backend="eager")
def f(x):
torch._dynamo.eval_frame._set_stance("force_eager")
return x + 1
with self.assertRaisesRegex(
AssertionError, "Attempt to trace forbidden callable"
):
f(torch.ones(3))
@torch.compile(backend="eager")
def g(x):
# cause a skipped frame
try:
torch._dynamo.graph_break()
except Exception:
pass
# NOTE: torch._dynamo.is_compiling() will get traced
# and return true. torch.compiler.is_compiling() is skipped
# and will return false.
if torch.compiler.is_compiling():
raise RuntimeError("Expect this frame to be skipped")
# should not be traced, but eval frame callback is still set
with torch.compiler.set_stance("force_eager"):
return x + 1
with self.assertRaisesRegex(RuntimeError, "set_stance in a torch.compile"):
g(torch.ones(3))
def test_set_stance_force_backend(self):
@torch.compile
def a(x):
return x + 1
cnts = torch._dynamo.testing.CompileCounter()
@torch.compiler.set_stance("default", force_backend=cnts)
def b(x):
return a(x)
b(torch.ones(3))
self.assertEqual(cnts.frame_count, 1)
@torch.compiler.set_stance("default", force_backend="eager")
def c(x):
return a(x)
# just make sure this doesn't crash
c(torch.ones(3))
with self.assertRaisesRegex(RuntimeError, "force_backend"):
@torch.compiler.set_stance("force_eager", force_backend="eager")
def d(x):
pass
def test_set_stance_force_backend_with_disable(self):
@torch.compiler.disable
def inner(x):
return x
@torch.compile(backend="eager")
def f(x):
return inner(x)
f(torch.randn(3, 3))
def fail_backend(gm, ex):
raise RuntimeError("fail!")
# should not raise error
with torch.compiler.set_stance("default", force_backend=fail_backend):
f(torch.randn(3, 3))
if __name__ == "__main__":
from torch._dynamo.test_case import run_tests
run_tests()
|