Skip to content
New issue

Have a question about this project? Sign up for a free GitHub account to open an issue and contact its maintainers and the community.

By clicking “Sign up for GitHub”, you agree to our terms of service and privacy statement. We’ll occasionally send you account related emails.

Already on GitHub? Sign in to your account

feat: support activation dynamo converters #2254

Merged
merged 3 commits into from
Sep 1, 2023
Merged
Show file tree
Hide file tree
Changes from all commits
Commits
File filter

Filter by extension

Filter by extension


Conversations
Failed to load comments.
Loading
Jump to
Jump to file
Failed to load files.
Loading
Diff view
Diff view
146 changes: 128 additions & 18 deletions py/torch_tensorrt/dynamo/conversion/aten_ops_converters.py
Original file line number Diff line number Diff line change
Expand Up @@ -152,15 +152,15 @@ def aten_ops_fmod(
return impl.elementwise.fmod(network, target, SourceIR.ATEN, name, args[0], args[1])


@dynamo_tensorrt_converter(torch.ops.aten.gelu.default) # type: ignore[misc]
def aten_ops_gelu(
@dynamo_tensorrt_converter(torch.ops.aten.relu.default)
def aten_ops_relu(
network: TRTNetwork,
target: Target,
args: Tuple[Argument, ...],
kwargs: Dict[str, Argument],
name: str,
) -> Union[TRTTensor, Sequence[TRTTensor]]:
return impl.activation.gelu(
return impl.activation.relu(
network,
target,
SourceIR.ATEN,
Expand All @@ -169,61 +169,171 @@ def aten_ops_gelu(
)


@dynamo_tensorrt_converter(torch.ops.aten.matmul) # type: ignore[misc]
@dynamo_tensorrt_converter(torch.ops.aten.mm.default) # type: ignore[misc]
@dynamo_tensorrt_converter(torch.ops.aten.mv.default) # type: ignore[misc]
def aten_ops_matmul(
@dynamo_tensorrt_converter(torch.ops.aten.sigmoid.default)
def aten_ops_sigmoid(
network: TRTNetwork,
target: Target,
args: Tuple[Argument, ...],
kwargs: Dict[str, Argument],
name: str,
) -> Union[TRTTensor, Sequence[TRTTensor]]:
return impl.matmul.matrix_multiply(
return impl.activation.sigmoid(
network,
target,
SourceIR.ATEN,
name,
args[0],
args[1],
)


@dynamo_tensorrt_converter(torch.ops.aten.layer_norm.default) # type: ignore[misc]
def aten_ops_layernorm(
@dynamo_tensorrt_converter(torch.ops.aten.tanh.default)
def aten_ops_tanh(
network: TRTNetwork,
target: Target,
args: Tuple[Argument, ...],
kwargs: Dict[str, Argument],
name: str,
) -> Union[TRTTensor, Sequence[TRTTensor]]:
return impl.normalization.layer_norm(
return impl.activation.tanh(
network,
target,
SourceIR.ATEN,
name,
args[0],
)


@dynamo_tensorrt_converter(torch.ops.aten.leaky_relu.default)
def aten_ops_leaky_relu(
network: TRTNetwork,
target: Target,
args: Tuple[Argument, ...],
kwargs: Dict[str, Argument],
name: str,
) -> Union[TRTTensor, Sequence[TRTTensor]]:
return impl.activation.leaky_relu(
network,
target,
SourceIR.ATEN,
name,
args[0],
args_bounds_check(args, 1, 0.01),
)


@dynamo_tensorrt_converter(torch.ops.aten.elu.default)
def aten_ops_elu(
network: TRTNetwork,
target: Target,
args: Tuple[Argument, ...],
kwargs: Dict[str, Argument],
name: str,
) -> Union[TRTTensor, Sequence[TRTTensor]]:
return impl.activation.elu(
network,
target,
SourceIR.ATEN,
name,
args[0],
alpha=args_bounds_check(args, 1, 1.0),
beta=args_bounds_check(args, 2, None),
)


@dynamo_tensorrt_converter(torch.ops.aten.softplus.default)
def aten_ops_softplus(
network: TRTNetwork,
target: Target,
args: Tuple[Argument, ...],
kwargs: Dict[str, Argument],
name: str,
) -> Union[TRTTensor, Sequence[TRTTensor]]:
return impl.activation.softplus(
network,
target,
SourceIR.ATEN,
name,
args[0],
beta=args_bounds_check(args, 1, 1),
)


@dynamo_tensorrt_converter(torch.ops.aten.clip.default)
def aten_ops_clip(
network: TRTNetwork,
target: Target,
args: Tuple[Argument, ...],
kwargs: Dict[str, Argument],
name: str,
) -> Union[TRTTensor, Sequence[TRTTensor]]:
return impl.activation.clip(
network,
target,
SourceIR.ATEN,
name,
args[0],
alpha=args_bounds_check(args, 1),
beta=args_bounds_check(args, 2),
)


@dynamo_tensorrt_converter(torch.ops.aten.hardsigmoid.default)
def aten_ops_hard_sigmoid(
network: TRTNetwork,
target: Target,
args: Tuple[Argument, ...],
kwargs: Dict[str, Argument],
name: str,
) -> Union[TRTTensor, Sequence[TRTTensor]]:
return impl.activation.hard_sigmoid(
network,
target,
SourceIR.ATEN,
name,
args[0],
alpha=args_bounds_check(args, 1, 1 / 6),
beta=args_bounds_check(args, 2, 1 / 2),
)


@dynamo_tensorrt_converter(torch.ops.aten.matmul) # type: ignore[misc]
@dynamo_tensorrt_converter(torch.ops.aten.mm.default) # type: ignore[misc]
@dynamo_tensorrt_converter(torch.ops.aten.mv.default) # type: ignore[misc]
def aten_ops_matmul(
network: TRTNetwork,
target: Target,
args: Tuple[Argument, ...],
kwargs: Dict[str, Argument],
name: str,
) -> Union[TRTTensor, Sequence[TRTTensor]]:
return impl.matmul.matrix_multiply(
network,
target,
SourceIR.ATEN,
name,
args[0],
args[1],
args[2],
args[3],
args[4],
)


@dynamo_tensorrt_converter(torch.ops.aten.relu.default) # type: ignore[misc]
def aten_ops_relu(
@dynamo_tensorrt_converter(torch.ops.aten.layer_norm.default) # type: ignore[misc]
def aten_ops_layernorm(
network: TRTNetwork,
target: Target,
args: Tuple[Argument, ...],
kwargs: Dict[str, Argument],
name: str,
) -> Union[TRTTensor, Sequence[TRTTensor]]:
return impl.activation.relu(
return impl.normalization.layer_norm(
network,
target,
SourceIR.ATEN,
name,
args[0],
args[1],
args[2],
args[3],
args[4],
)


Expand Down
63 changes: 0 additions & 63 deletions py/torch_tensorrt/dynamo/conversion/impl/activation.py

This file was deleted.

Original file line number Diff line number Diff line change
@@ -0,0 +1 @@
from .ops import *
42 changes: 42 additions & 0 deletions py/torch_tensorrt/dynamo/conversion/impl/activation/base.py
Original file line number Diff line number Diff line change
@@ -0,0 +1,42 @@
from typing import Any, Callable, Optional

import tensorrt as trt
from torch.fx.node import Target
from torch_tensorrt.dynamo._SourceIR import SourceIR
from torch_tensorrt.fx.converters.converter_utils import (
mark_as_int8_layer,
set_layer_name,
)
from torch_tensorrt.fx.types import TRTNetwork, TRTTensor


def convert_activation(
network: TRTNetwork,
target: Target,
source_ir: Optional[SourceIR],
name: str,
operation_type: trt.ActivationType,
input_val: TRTTensor,
alpha: Optional[Any] = None,
beta: Optional[Any] = None,
dyn_range_fn: Optional[Callable[[float, float], Any]] = None,
) -> TRTTensor:
"""
Add a TensorRT Activation layer to `network`.
"""
if not isinstance(input_val, TRTTensor):
raise RuntimeError(
f"{operation_type} received input {input_val} that is not part "
"of the TensorRT region!"
)
layer = network.add_activation(input_val, operation_type)
if alpha is not None:
layer.alpha = alpha
if beta is not None:
layer.beta = beta
set_layer_name(layer, target, name, source_ir)

if input_val.dynamic_range is not None:
dyn_range = dyn_range_fn(input_val.dynamic_range)
mark_as_int8_layer(layer, dyn_range)
return layer.get_output(0)
Loading