Skip to content

Commit 755462a

Browse files
authored
Additional converters for floordiv, mod, ne, and torch::tensor() operations (#505)
* Initioal version of ne, floordiv, mod and tensor converters. Extend ops for relu and sigmoid. * Converters for floordiv, mod, ne, and torch::tensor() operations . Extend relu and sigmoid converters to Tensor methods. * Update CHANGELOG.md
1 parent c294dac commit 755462a

File tree

8 files changed

+293
-2
lines changed

8 files changed

+293
-2
lines changed

CHANGELOG.md

Lines changed: 3 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -10,3 +10,6 @@
1010
- Added GroupNorm plugin which internally uses PyTorch aten::group_norm
1111
- Replaced Tensor.ndim references with len(tensor.shape) to support older pytorch versions
1212
- Added reduced precision documentation page
13+
- Added converters for ``floordiv``, ``mod``, ``ne``, and ``torch.tensor`` operations
14+
- Extended ``relu`` converter to support ``Tensor.relu`` operation
15+
- Extended ``sigmoid`` converter to support ``Tensor.sigmoid`` operation

torch2trt/converters/__init__.py

Lines changed: 4 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -26,6 +26,7 @@
2626
from .compare import *
2727
from .div import *
2828
from .expand import *
29+
from .floordiv import *
2930
from .getitem import *
3031
from .identity import *
3132
from .instance_norm import *
@@ -35,8 +36,10 @@
3536
from .max_pool2d import *
3637
from .mean import *
3738
from .min import *
39+
from .mod import *
3840
from .mul import *
3941
from .normalize import *
42+
from .ne import *
4043
from .narrow import *
4144
from .pad import *
4245
from .permute import *
@@ -52,6 +55,7 @@
5255
from .sub import *
5356
from .sum import *
5457
from .tanh import *
58+
from .tensor import *
5559
from .transpose import *
5660
from .unary import *
5761
from .view import *

torch2trt/converters/floordiv.py

Lines changed: 81 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -0,0 +1,81 @@
1+
from torch2trt.torch2trt import *
2+
from torch2trt.module_test import add_module_test
3+
4+
5+
@tensorrt_converter('torch.Tensor.__floordiv__')
6+
@tensorrt_converter('torch.Tensor.__ifloordiv__')
7+
@tensorrt_converter('torch.floor_divide')
8+
def convert_floordiv(ctx):
9+
input_a = ctx.method_args[0]
10+
input_b = ctx.method_args[1]
11+
output = ctx.method_return
12+
input_a_trt, input_b_trt = add_missing_trt_tensors(ctx.network, [input_a, input_b])
13+
input_a_trt, input_b_trt = broadcast_trt_tensors(ctx.network, [input_a_trt, input_b_trt], len(output.shape) - 1)
14+
# we can not use ElementWiseOperation.FLOOR_DIV directly because Torch truncate negative result toward 0
15+
# but TensorRT FLOOR_DIV op toward -Inf
16+
# sign = ab / |ab|
17+
# floordiv result: sign * (|a| // |b|)
18+
ab_layer = ctx.network.add_elementwise(input_a_trt, input_b_trt, trt.ElementWiseOperation.PROD)
19+
abs_ab_layer = ctx.network.add_unary(ab_layer.get_output(0), trt.UnaryOperation.ABS)
20+
sign_layer = ctx.network.add_elementwise(ab_layer.get_output(0), abs_ab_layer.get_output(0),
21+
trt.ElementWiseOperation.DIV)
22+
abs_a_layer = ctx.network.add_unary(input_a_trt, trt.UnaryOperation.ABS)
23+
abs_b_layer = ctx.network.add_unary(input_b_trt, trt.UnaryOperation.ABS)
24+
abs_floor_layer = ctx.network.add_elementwise(abs_a_layer.get_output(0), abs_b_layer.get_output(0),
25+
trt.ElementWiseOperation.FLOOR_DIV)
26+
out_layer = ctx.network.add_elementwise(sign_layer.get_output(0), abs_floor_layer.get_output(0),
27+
trt.ElementWiseOperation.PROD)
28+
output._trt = out_layer.get_output(0)
29+
30+
31+
class FloorDiv(torch.nn.Module):
32+
def __init__(self):
33+
super(FloorDiv, self).__init__()
34+
35+
def forward(self, x, y):
36+
return x // y
37+
38+
39+
@add_module_test(torch.float32, torch.device('cuda'), [(1, 3, 40, 20), (1, 3, 1, 20)])
40+
def test_floordiv_op():
41+
return FloorDiv()
42+
43+
44+
class FloorDivAssign (torch.nn.Module):
45+
def __init__(self):
46+
super(FloorDivAssign, self).__init__()
47+
48+
def forward(self, x, y):
49+
x //= y
50+
return x
51+
52+
53+
@add_module_test(torch.float32, torch.device('cuda'), [(1, 3, 40, 20), (1, 3, 1, 20)])
54+
def test_floordiv_op_assign():
55+
return FloorDivAssign()
56+
57+
58+
class FloorDivConst(torch.nn.Module):
59+
def __init__(self):
60+
super(FloorDivConst, self).__init__()
61+
62+
def forward(self, x):
63+
return x // 2.
64+
65+
66+
@add_module_test(torch.float32, torch.device('cuda'), [(1, 3, 40, 20)])
67+
def test_floordiv_op_const():
68+
return FloorDivConst()
69+
70+
71+
class TorchFloorDiv(torch.nn.Module):
72+
def __init__(self):
73+
super(TorchFloorDiv, self).__init__()
74+
75+
def forward(self, x, y):
76+
return torch.floor_divide(x, y)
77+
78+
79+
@add_module_test(torch.float32, torch.device('cuda'), [(1, 3, 40, 20), (1, 3, 1, 20)])
80+
def test_floordiv_func():
81+
return TorchFloorDiv()

torch2trt/converters/mod.py

Lines changed: 99 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -0,0 +1,99 @@
1+
from torch2trt.torch2trt import *
2+
from torch2trt.module_test import add_module_test
3+
4+
5+
@tensorrt_converter('torch.fmod')
6+
def convert_mod(ctx):
7+
input_a = ctx.method_args[0]
8+
input_b = ctx.method_args[1]
9+
output = ctx.method_return
10+
input_a_trt, input_b_trt = add_missing_trt_tensors(ctx.network, [input_a, input_b])
11+
input_a_trt, input_b_trt = broadcast_trt_tensors(ctx.network, [input_a_trt, input_b_trt], len(output.shape) - 1)
12+
# we can not use ElementWiseOperation.FLOOR_DIV directly because Torch truncate negative result toward 0
13+
# but TensorRT FLOOR_DIV op toward -Inf
14+
# sign = ab / |ab|
15+
# floordiv result: sign * (|a| // |b|)
16+
ab_layer = ctx.network.add_elementwise(input_a_trt, input_b_trt, trt.ElementWiseOperation.PROD)
17+
abs_ab_layer = ctx.network.add_unary(ab_layer.get_output(0), trt.UnaryOperation.ABS)
18+
sign_layer = ctx.network.add_elementwise(ab_layer.get_output(0), abs_ab_layer.get_output(0),
19+
trt.ElementWiseOperation.DIV)
20+
abs_a_layer = ctx.network.add_unary(input_a_trt, trt.UnaryOperation.ABS)
21+
abs_b_layer = ctx.network.add_unary(input_b_trt, trt.UnaryOperation.ABS)
22+
abs_floor_layer = ctx.network.add_elementwise(abs_a_layer.get_output(0), abs_b_layer.get_output(0),
23+
trt.ElementWiseOperation.FLOOR_DIV)
24+
# a % b = a - (a//b) * b
25+
floordiv_layer = ctx.network.add_elementwise(sign_layer.get_output(0), abs_floor_layer.get_output(0),
26+
trt.ElementWiseOperation.PROD)
27+
prod_layer = ctx.network.add_elementwise(floordiv_layer.get_output(0), input_b_trt, trt.ElementWiseOperation.PROD)
28+
sub_layer = ctx.network.add_elementwise(input_a_trt, prod_layer.get_output(0), trt.ElementWiseOperation.SUB)
29+
output._trt = sub_layer.get_output(0)
30+
31+
32+
@tensorrt_converter('torch.Tensor.__mod__')
33+
# we need separate converter for operator because for some reason Torch use truncation toward -Inf for this op.
34+
# bug is filed: https://github.com/pytorch/pytorch/issues/52425
35+
# but for now we have to convert model exactly
36+
def convert_mod(ctx):
37+
input_a = ctx.method_args[0]
38+
input_b = ctx.method_args[1]
39+
output = ctx.method_return
40+
input_a_trt, input_b_trt = add_missing_trt_tensors(ctx.network, [input_a, input_b])
41+
input_a_trt, input_b_trt = broadcast_trt_tensors(ctx.network, [input_a_trt, input_b_trt], len(output.shape) - 1)
42+
# a % b = a - (a//b) * b
43+
floordiv_layer = ctx.network.add_elementwise(input_a_trt, input_b_trt, trt.ElementWiseOperation.FLOOR_DIV)
44+
prod_layer = ctx.network.add_elementwise(floordiv_layer.get_output(0), input_b_trt, trt.ElementWiseOperation.PROD)
45+
mod_layer = ctx.network.add_elementwise(input_a_trt, prod_layer.get_output(0), trt.ElementWiseOperation.SUB)
46+
output._trt = mod_layer.get_output(0)
47+
48+
49+
class Mod(torch.nn.Module):
50+
def __init__(self):
51+
super(Mod, self).__init__()
52+
53+
def forward(self, x, y):
54+
return x % y
55+
56+
57+
@add_module_test(torch.float32, torch.device('cuda'), [(1, 3, 40, 20), (1, 3, 1, 20)])
58+
def test_mod_op():
59+
return Mod()
60+
61+
62+
class ModAssign(torch.nn.Module):
63+
def __init__(self):
64+
super(ModAssign, self).__init__()
65+
66+
def forward(self, x, y):
67+
x %= y
68+
return x
69+
70+
71+
@add_module_test(torch.float32, torch.device('cuda'), [(1, 3, 40, 20), (1, 3, 1, 20)])
72+
def test_mod_op_assign():
73+
return ModAssign()
74+
75+
76+
class ModConst(torch.nn.Module):
77+
def __init__(self):
78+
super(ModConst, self).__init__()
79+
80+
def forward(self, x):
81+
return x % 2.
82+
83+
84+
@add_module_test(torch.float32, torch.device('cuda'), [(1, 3, 40, 20)])
85+
def test_mod_op_const():
86+
return ModConst()
87+
88+
89+
class TorchMod(torch.nn.Module):
90+
def __init__(self):
91+
super(TorchMod, self).__init__()
92+
93+
def forward(self, x, y):
94+
return torch.fmod(x, y)
95+
96+
97+
@add_module_test(torch.float32, torch.device('cuda'), [(1, 3, 40, 20), (1, 3, 40, 20)])
98+
def test_mod_func():
99+
return TorchMod()

torch2trt/converters/ne.py

Lines changed: 54 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -0,0 +1,54 @@
1+
from torch2trt.torch2trt import *
2+
from torch2trt.module_test import add_module_test
3+
4+
5+
@tensorrt_converter('torch.ne')
6+
@tensorrt_converter('torch.Tensor.__ne__')
7+
def convert_ne(ctx):
8+
input_a = ctx.method_args[0]
9+
input_b = ctx.method_args[1]
10+
output = ctx.method_return
11+
input_a_trt, input_b_trt = add_missing_trt_tensors(ctx.network, [input_a, input_b])
12+
input_a_trt, input_b_trt = broadcast_trt_tensors(ctx.network, [input_a_trt, input_b_trt], len(output.shape) - 1)
13+
layer_1 = ctx.network.add_elementwise(input_a_trt, input_b_trt, trt.ElementWiseOperation.EQUAL)
14+
layer_2 = ctx.network.add_unary(layer_1.get_output(0), trt.UnaryOperation.NOT)
15+
output._trt = layer_2.get_output(0)
16+
17+
18+
class NotEqual(torch.nn.Module):
19+
def __init__(self):
20+
super(NotEqual, self).__init__()
21+
22+
def forward(self, x, y):
23+
return x != y
24+
25+
26+
@add_module_test(torch.float32, torch.device('cuda'), [(1, 3, 40, 20), (1, 3, 1, 20)])
27+
def test_ne_op():
28+
return NotEqual()
29+
30+
31+
class NotEqualConst(torch.nn.Module):
32+
def __init__(self):
33+
super(NotEqualConst, self).__init__()
34+
35+
def forward(self, x):
36+
return x != 13.62
37+
38+
39+
@add_module_test(torch.float32, torch.device('cuda'), [(1, 3, 40, 20)])
40+
def test_ne_op_const():
41+
return NotEqualConst()
42+
43+
44+
class TorchNotEqual(torch.nn.Module):
45+
def __init__(self):
46+
super(TorchNotEqual, self).__init__()
47+
48+
def forward(self, x, y):
49+
return torch.ne(x, y)
50+
51+
52+
@add_module_test(torch.float32, torch.device('cuda'), [(1, 3, 40, 20), (1, 3, 1, 20)])
53+
def test_ne_torch():
54+
return TorchNotEqual()

torch2trt/converters/relu.py

Lines changed: 15 additions & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -6,6 +6,7 @@
66
@tensorrt_converter('torch.relu_')
77
@tensorrt_converter('torch.nn.functional.relu')
88
@tensorrt_converter('torch.nn.functional.relu_')
9+
@tensorrt_converter('torch.Tensor.relu')
910
def convert_functional_relu(ctx):
1011
ctx.method_args = (torch.nn.ReLU(),) + ctx.method_args
1112
convert_relu(ctx)
@@ -32,4 +33,17 @@ def forward(self, x):
3233

3334
@add_module_test(torch.float32, torch.device('cuda'), [(1, 3, 4, 5)])
3435
def test_functional_relu_basic():
35-
return FunctionalRelu()
36+
return FunctionalRelu()
37+
38+
39+
class TensorRelu(torch.nn.Module):
40+
def __init__(self):
41+
super(TensorRelu, self).__init__()
42+
43+
def forward(self, x):
44+
return x.relu()
45+
46+
47+
@add_module_test(torch.float32, torch.device('cuda'), [(1, 3, 40, 20)])
48+
def test_tensor_relu():
49+
return TensorRelu()

torch2trt/converters/sigmoid.py

Lines changed: 15 additions & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -4,6 +4,7 @@
44

55
@tensorrt_converter('torch.nn.functional.sigmoid')
66
@tensorrt_converter('torch.sigmoid')
7+
@tensorrt_converter('torch.Tensor.sigmoid')
78
def convert_sigmoid(ctx):
89
input = ctx.method_args[0]
910
input_trt = add_missing_trt_tensors(ctx.network, [input])[0]
@@ -15,4 +16,17 @@ def convert_sigmoid(ctx):
1516

1617
@add_module_test(torch.float32, torch.device('cuda'), [(1, 3, 3, 3)])
1718
def test_sigmoid_basic():
18-
return torch.nn.Sigmoid()
19+
return torch.nn.Sigmoid()
20+
21+
22+
class TensorSigmoid(torch.nn.Module):
23+
def __init__(self):
24+
super(TensorSigmoid, self).__init__()
25+
26+
def forward(self, x):
27+
return x.sigmoid()
28+
29+
30+
@add_module_test(torch.float32, torch.device('cuda'), [(1, 3, 40, 20)])
31+
def test_tensor_sigmoid():
32+
return TensorSigmoid()

torch2trt/converters/tensor.py

Lines changed: 22 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -0,0 +1,22 @@
1+
from torch2trt.torch2trt import *
2+
from torch2trt.module_test import add_module_test
3+
4+
5+
@tensorrt_converter('torch.tensor')
6+
def convert_mod(ctx):
7+
output = ctx.method_return
8+
layer = ctx.network.add_constant(tuple(output.shape), output.detach().cpu().numpy() )
9+
output._trt = layer.get_output(0)
10+
11+
12+
class TorchTensor(torch.nn.Module):
13+
def __init__(self):
14+
super(TorchTensor, self).__init__()
15+
16+
def forward(self, x):
17+
return x + torch.tensor([[1., 2., 3.], [4., 5., 6.]], device=torch.device('cuda'))
18+
19+
20+
@add_module_test(torch.float32, torch.device('cuda'), [(1, 2, 3)])
21+
def test_tensor_creation():
22+
return TorchTensor()

0 commit comments

Comments
 (0)