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

Adding Lora implementation for nn.Conv1d #2333

Open
wants to merge 5 commits into
base: main
Choose a base branch
from
Open
Changes from 1 commit
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
12 changes: 12 additions & 0 deletions src/peft/tuners/lora/layer.py
Original file line number Diff line number Diff line change
Expand Up @@ -67,6 +67,8 @@ def __init__(self, base_layer: nn.Module, ephemeral_gpu_offload: bool = False, *
base_layer = self.get_base_layer()
if isinstance(base_layer, nn.Linear):
in_features, out_features = base_layer.in_features, base_layer.out_features
elif isinstance(base_layer, nn.Conv1d):
in_features, out_features = base_layer.in_channels, base_layer.out_channels
elif isinstance(base_layer, nn.Conv2d):
in_features, out_features = base_layer.in_channels, base_layer.out_channels
elif isinstance(base_layer, nn.Conv3d):
Expand Down Expand Up @@ -1296,6 +1298,13 @@ def __init__(self, *args, **kwargs):
def _get_dora_layer_class(self):
return DoraConv2dLayer

class Conv1d(_ConvNd):
# Lora implemented in a conv3d layer
CCLDArjun marked this conversation as resolved.
Show resolved Hide resolved
def __init__(self, *args, **kwargs):
super().__init__(*args, **kwargs)
if not self._kernel_dim == 3:
raise ValueError(f"Conv1d layer kernel must have 3 dimensions, not {self._kernel_dim}")
self.conv_fn = F.conv1d
CCLDArjun marked this conversation as resolved.
Show resolved Hide resolved

class Conv3d(_ConvNd):
# Lora implemented in a conv3d layer
Expand Down Expand Up @@ -1679,6 +1688,9 @@ def dispatch_default(
elif isinstance(target_base_layer, torch.nn.Conv3d):
kwargs.update(lora_config.loftq_config)
new_module = Conv3d(target, adapter_name, **kwargs)
elif isinstance(target_base_layer, nn.Conv1d):
kwargs.update(lora_config.loftq_config)
new_module = Conv1d(target, adapter_name, **kwargs)
elif isinstance(target_base_layer, torch.nn.MultiheadAttention):
kwargs.update(lora_config.loftq_config)
new_module = MultiheadAttention(target, adapter_name, **kwargs)
Expand Down